Skip to content

Commit

Permalink
Removed legacy keras import
Browse files Browse the repository at this point in the history
  • Loading branch information
RobertSamoilescu committed Dec 5, 2024
1 parent 77a0a46 commit a39d198
Show file tree
Hide file tree
Showing 21 changed files with 254 additions and 345 deletions.
4 changes: 2 additions & 2 deletions Makefile
Original file line number Diff line number Diff line change
Expand Up @@ -67,5 +67,5 @@ repl:

.PHONY: test
test:
pytest alibi/explainers/tests_tf1/
pytest alibi --ignore=alibi/explainers/tests_tf1/
TF_USE_LEGACY_KERAS=1 pytest -m "tf1" alibi
pytest -m "not tf1" alibi
5 changes: 3 additions & 2 deletions alibi/datasets/tensorflow.py
Original file line number Diff line number Diff line change
@@ -1,7 +1,8 @@
import numpy as np
from typing import Tuple, Union

from alibi.utils.legacy_keras import keras
import tensorflow.keras as keras
import numpy as np

from alibi.utils.data import Bunch


Expand Down
2 changes: 1 addition & 1 deletion alibi/explainers/backends/tensorflow/cfrl_base.py
Original file line number Diff line number Diff line change
Expand Up @@ -9,8 +9,8 @@

import numpy as np
import tensorflow as tf
import tensorflow.keras as keras

from alibi.utils.legacy_keras import keras
from alibi.explainers.backends.cfrl_base import CounterfactualRLDataset
from alibi.models.tensorflow.actor_critic import Actor, Critic

Expand Down
45 changes: 22 additions & 23 deletions alibi/explainers/integrated_gradients.py
Original file line number Diff line number Diff line change
Expand Up @@ -8,7 +8,6 @@
import numpy as np
import tensorflow as tf

from alibi.utils.legacy_keras import keras
from alibi.api.defaults import DEFAULT_DATA_INTGRAD, DEFAULT_META_INTGRAD
from alibi.api.interfaces import Explainer, Explanation
from alibi.utils.approximation_methods import approximation_parameters
Expand All @@ -18,7 +17,7 @@
_valid_output_shape_type: List = [tuple, list]


def _compute_convergence_delta(model: Union[keras.models.Model],
def _compute_convergence_delta(model: Union[tf.keras.models.Model],
input_dtypes: List[tf.DType],
attributions: List[np.ndarray],
start_point: Union[List[np.ndarray], np.ndarray],
Expand Down Expand Up @@ -128,7 +127,7 @@ def _select_target(preds: tf.Tensor,
return preds


def _run_forward(model: Union[keras.models.Model],
def _run_forward(model: Union[tf.keras.models.Model],
x: Union[List[tf.Tensor], List[np.ndarray], tf.Tensor, np.ndarray],
target: Union[None, tf.Tensor, np.ndarray, list],
forward_kwargs: Optional[dict] = None) -> tf.Tensor:
Expand Down Expand Up @@ -162,8 +161,8 @@ def _run_forward(model: Union[keras.models.Model],
return preds


def _run_forward_from_layer(model: keras.models.Model,
layer: keras.layers.Layer,
def _run_forward_from_layer(model: tf.keras.models.Model,
layer: tf.keras.layers.Layer,
orig_call: Callable,
orig_dummy_input: Union[list, np.ndarray],
x: tf.Tensor,
Expand Down Expand Up @@ -242,8 +241,8 @@ def wrapper(*args, **kwargs):
return preds


def _run_forward_to_layer(model: keras.models.Model,
layer: keras.layers.Layer,
def _run_forward_to_layer(model: tf.keras.models.Model,
layer: tf.keras.layers.Layer,
orig_call: Callable,
x: Union[List[np.ndarray], np.ndarray],
forward_kwargs: Optional[dict] = None,
Expand Down Expand Up @@ -310,8 +309,8 @@ def wrapper(*args, **kwargs):

def _forward_input_baseline(X: Union[List[np.ndarray], np.ndarray],
bls: Union[List[np.ndarray], np.ndarray],
model: keras.Model,
layer: keras.layers.Layer,
model: tf.keras.Model,
layer: tf.keras.layers.Layer,
orig_call: Callable,
forward_kwargs: Optional[dict] = None,
forward_to_inputs: bool = False) -> Tuple[Union[list, tf.Tensor], Union[list, tf.Tensor]]:
Expand Down Expand Up @@ -369,7 +368,7 @@ def _forward_input_baseline(X: Union[List[np.ndarray], np.ndarray],
return X, bls


def _gradients_input(model: Union[keras.models.Model],
def _gradients_input(model: Union[tf.keras.models.Model],
x: List[tf.Tensor],
target: Union[None, tf.Tensor],
forward_kwargs: Optional[dict] = None) -> List[tf.Tensor]:
Expand Down Expand Up @@ -404,8 +403,8 @@ def _gradients_input(model: Union[keras.models.Model],
return grads


def _gradients_layer(model: Union[keras.models.Model],
layer: Union[keras.layers.Layer],
def _gradients_layer(model: Union[tf.keras.models.Model],
layer: Union[tf.keras.layers.Layer],
orig_call: Callable,
orig_dummy_input: Union[list, np.ndarray],
x: Union[List[tf.Tensor], tf.Tensor],
Expand Down Expand Up @@ -631,7 +630,7 @@ def _check_target(output_shape: Tuple,


def _get_target_from_target_fn(target_fn: Callable,
model: keras.Model,
model: tf.keras.Model,
X: Union[np.ndarray, List[np.ndarray]],
forward_kwargs: Optional[dict] = None) -> np.ndarray:
"""
Expand Down Expand Up @@ -705,7 +704,7 @@ def _sum_integral_terms(step_sizes: list,


def _calculate_sum_int(batches: List[List[tf.Tensor]],
model: Union[keras.Model],
model: Union[tf.keras.Model],
target: Optional[np.ndarray],
target_paths: np.ndarray,
n_steps: int,
Expand Down Expand Up @@ -755,7 +754,7 @@ def _calculate_sum_int(batches: List[List[tf.Tensor]],
return sum_int


def _validate_output(model: keras.Model,
def _validate_output(model: tf.keras.Model,
target: Optional[np.ndarray]) -> None:
"""
Validates the model's output type and raises an error if the output type is not supported.
Expand Down Expand Up @@ -790,11 +789,11 @@ class LayerState(str, Enum):
class IntegratedGradients(Explainer):

def __init__(self,
model: keras.Model,
model: tf.keras.Model,
layer: Optional[
Union[
Callable[[keras.Model], keras.layers.Layer],
keras.layers.Layer
Callable[[tf.keras.Model], tf.keras.layers.Layer],
tf.keras.layers.Layer
]
] = None,
target_fn: Optional[Callable] = None,
Expand Down Expand Up @@ -840,14 +839,14 @@ def __init__(self,
if self.model.inputs is None:
self._has_inputs = False
else:
self._has_inputs = False
self._has_inputs = True

if layer is None:
self.orig_call: Optional[Callable] = None
self.layer = None
layer_meta: Union[int, str] = LayerState.UNSPECIFIED.value

elif isinstance(layer, keras.layers.Layer):
elif isinstance(layer, tf.keras.layers.Layer):
self.orig_call = layer.call
self.layer = layer

Expand Down Expand Up @@ -988,7 +987,7 @@ def explain(self,
if not self._has_inputs:
# Inferring model's inputs from data points for models with no explicit inputs
# (typically subclassed models)
inputs = [keras.Input(shape=xx.shape[1:], dtype=xx.dtype) for xx in X]
inputs = [tf.keras.Input(shape=xx.shape[1:], dtype=xx.dtype) for xx in X]
self.model(inputs, **forward_kwargs)

_validate_output(self.model, target)
Expand Down Expand Up @@ -1036,7 +1035,7 @@ def explain(self,
else:
# Attributions calculation in case of single input
if not self._has_inputs:
inputs = keras.Input(shape=X.shape[1:], dtype=X.dtype) # type: ignore
inputs = tf.keras.Input(shape=X.shape[1:], dtype=X.dtype) # type: ignore
self.model(inputs, **forward_kwargs)

_validate_output(self.model, target)
Expand Down Expand Up @@ -1121,7 +1120,7 @@ def _build_explanation(self,

return Explanation(meta=copy.deepcopy(self.meta), data=data)

def reset_predictor(self, predictor: Union[keras.Model]) -> None:
def reset_predictor(self, predictor: Union[tf.keras.Model]) -> None:
"""
Resets the predictor model.
Expand Down
4 changes: 2 additions & 2 deletions alibi/explainers/tests/test_anchor_image.py
Original file line number Diff line number Diff line change
Expand Up @@ -2,8 +2,8 @@
from pytest_lazyfixture import lazy_fixture
import torch
import numpy as np
import tensorflow as tf

from alibi.utils.legacy_keras import keras
from alibi.api.defaults import DEFAULT_META_ANCHOR, DEFAULT_DATA_ANCHOR_IMG
from alibi.exceptions import PredictorCallError, PredictorReturnTypeError
from alibi.explainers.anchors.anchor_image import AnchorImage, AnchorImageSampler, scale_image
Expand All @@ -27,7 +27,7 @@ def predict_fn(request):
This fixture takes in a white-box model (Tensorflow or Pytorch) and returns an
AnchorImage compatible prediction function.
"""
if isinstance(request.param[0], keras.Model):
if isinstance(request.param[0], tf.keras.Model):
func = request.param[0].predict
elif isinstance(request.param[0], torch.nn.Module):
def func(image: np.ndarray) -> np.ndarray:
Expand Down
Original file line number Diff line number Diff line change
@@ -1,9 +1,13 @@
from alibi.api.defaults import DEFAULT_META_CEM, DEFAULT_DATA_CEM
from alibi.explainers import CEM
import numpy as np
from sklearn.datasets import load_iris
from sklearn.linear_model import LogisticRegression
import pytest


def test_cem(set_env_variables):
@pytest.mark.tf1
def test_cem(disable_tf2):
# load iris dataset
dataset = load_iris()

Expand All @@ -28,8 +32,6 @@ def test_cem(set_env_variables):
# test explainer initialization
shape = (1, 4)
feature_range = (X.min(axis=0).reshape(shape) - .1, X.max(axis=0).reshape(shape) + .1)

from alibi.explainers import CEM
cem = CEM(predict_fn, 'PN', shape, feature_range=feature_range, max_iterations=10, no_info_val=-1.)
explanation = cem.explain(X_expl, verbose=False)

Expand All @@ -38,7 +40,5 @@ def test_cem(set_env_variables):
assert (explanation.X != explanation.PN).astype(int).sum() > 0
assert explanation.X_pred != explanation.PN_pred
assert explanation.grads_graph.shape == explanation.grads_num.shape

from alibi.api.defaults import DEFAULT_META_CEM, DEFAULT_DATA_CEM
assert explanation.meta.keys() == DEFAULT_META_CEM.keys()
assert explanation.data.keys() == DEFAULT_DATA_CEM.keys()
150 changes: 150 additions & 0 deletions alibi/explainers/tests/test_cfproto.py
Original file line number Diff line number Diff line change
@@ -0,0 +1,150 @@
import numpy as np
import pytest
import tensorflow as tf
from alibi.api.defaults import DEFAULT_META_CFP, DEFAULT_DATA_CFP
from alibi.explainers import CounterfactualProto
from alibi.utils.mapping import ohe_to_ord, ord_to_num


@pytest.fixture
def tf_keras_iris_explainer(request, models, iris_data):
X_train = iris_data['X_train']
model, ae, enc = models
if request.param[0]: # use k-d trees
ae = None
enc = None

shape = (1, 4)
cf_explainer = CounterfactualProto(model, shape, gamma=100, theta=100,
ae_model=ae, enc_model=enc, use_kdtree=request.param[0],
max_iterations=1000, c_init=request.param[1], c_steps=request.param[2],
feature_range=(X_train.min(axis=0).reshape(shape),
X_train.max(axis=0).reshape(shape)))
yield model, cf_explainer
tf.keras.backend.clear_session()


@pytest.mark.tf1
@pytest.mark.parametrize('tf_keras_iris_explainer, use_kdtree, k',
[((False, 0., 1), False, None),
((False, 1., 3), False, None),
((False, 0., 1), False, 2),
((False, 1., 3), False, 2),
((True, 0., 1), True, None),
((True, 1., 3), True, None),
((True, 0., 1), True, 2),
((True, 1., 3), True, 2)],
indirect=['tf_keras_iris_explainer'])
@pytest.mark.parametrize('models',
[('iris-ffn-tf2.2.0', 'iris-ae-tf2.2.0', 'iris-enc-tf2.2.0'),
('iris-ffn-tf1.15.2.h5', 'iris-ae-tf1.15.2.h5', 'iris-enc-tf1.15.2.h5')],
indirect=True)
def test_tf_keras_iris_explainer(disable_tf2, iris_data, tf_keras_iris_explainer, use_kdtree, k):
model, cf = tf_keras_iris_explainer
X_train = iris_data['X_train']

# instance to be explained
x = X_train[0].reshape(1, -1)
pred_class = np.argmax(model.predict(x))
not_pred_class = np.argmin(model.predict(x))

# test fit
cf.fit(X_train)
if use_kdtree: # k-d trees
assert len(cf.kdtrees) == cf.classes # each class has a k-d tree
n_by_class = 0
for c in range(cf.classes):
n_by_class += cf.X_by_class[c].shape[0]
assert n_by_class == X_train.shape[0] # all training instances are stored in the trees
assert cf.kdtrees[pred_class].query(x, k=1)[0] == 0. # nearest distance to own class equals 0
assert cf.score(x, not_pred_class, pred_class) == 0. # test score fn
else: # encoder
assert len(list(cf.class_proto.keys())) == cf.classes
assert [True for _ in range(cf.classes)] == [v.shape == (1, 2) for _, v in cf.class_proto.items()]
n_by_class = 0
for c in range(cf.classes):
n_by_class += cf.class_enc[c].shape[0]
assert n_by_class == X_train.shape[0] # all training instances encoded

# test explanation
explanation = cf.explain(x, k=k)
assert cf.id_proto != pred_class
assert np.argmax(model.predict(explanation.cf['X'])) == explanation.cf['class']
assert explanation.cf['grads_num'].shape == explanation.cf['grads_graph'].shape == x.shape
assert explanation.meta.keys() == DEFAULT_META_CFP.keys()
assert explanation.data.keys() == DEFAULT_DATA_CFP.keys()

# test gradient shapes
y = np.zeros((1, cf.classes))
np.put(y, pred_class, 1)
cf.predict = cf.predict.predict # make model black box
grads = cf.get_gradients(x, y, x.shape[1:], cf.cat_vars_ord)
assert grads.shape == x.shape


@pytest.fixture
def tf_keras_adult_explainer(request, models, adult_data):
shape = (1, 57)
cat_vars_ohe = adult_data['metadata']['cat_vars_ohe']
cf_explainer = CounterfactualProto(models[0], shape, beta=.01, cat_vars=cat_vars_ohe, ohe=True,
use_kdtree=request.param[0], max_iterations=1000,
c_init=request.param[1], c_steps=request.param[2],
feature_range=(-1 * np.ones((1, 12)), np.ones((1, 12))))
yield models[0], cf_explainer
tf.keras.backend.clear_session()


@pytest.mark.tf1
@pytest.mark.parametrize('tf_keras_adult_explainer, use_kdtree, k, d_type',
[((False, 1., 3), False, None, 'mvdm'),
((True, 1., 3), True, 2, 'mvdm'),
((True, 1., 3), True, 2, 'abdm')],
indirect=['tf_keras_adult_explainer'])
@pytest.mark.parametrize('models',
[('adult-ffn-tf2.2.0',), ('adult-ffn-tf1.15.2.h5',)],
ids='model={}'.format,
indirect=True)
def test_tf_keras_adult_explainer(disable_tf2, adult_data, tf_keras_adult_explainer, use_kdtree, k, d_type):
model, cf = tf_keras_adult_explainer
X_train = adult_data['preprocessor'].transform(adult_data['X_train']).toarray()

# instance to be explained
x = X_train[0].reshape(1, -1)
pred_class = np.argmax(model.predict(x))

# test fit
cf.fit(X_train, d_type=d_type)

# checked ragged tensor shape
n_cat = len(list(cf.cat_vars_ord.keys()))
max_key = max(cf.cat_vars_ord, key=cf.cat_vars_ord.get)
max_cat = cf.cat_vars_ord[max_key]
assert cf.d_abs_ragged.shape == (n_cat, max_cat)

if use_kdtree: # k-d trees
assert len(cf.kdtrees) == cf.classes # each class has a k-d tree
n_by_class = 0
for c in range(cf.classes):
n_by_class += cf.X_by_class[c].shape[0]
assert n_by_class == X_train.shape[0] # all training instances are stored in the trees

# test explanation
explanation = cf.explain(x, k=k)
if use_kdtree:
assert cf.id_proto != pred_class
assert np.argmax(model.predict(explanation.cf['X'])) == explanation.cf['class']
num_shape = (1, 12)
assert explanation.cf['grads_num'].shape == explanation.cf['grads_graph'].shape == num_shape
assert explanation.meta.keys() == DEFAULT_META_CFP.keys()
assert explanation.data.keys() == DEFAULT_DATA_CFP.keys()

# test gradient shapes
y = np.zeros((1, cf.classes))
np.put(y, pred_class, 1)
cf.predict = cf.predict.predict # make model black box
# convert instance to numerical space
x_ord = ohe_to_ord(x, cf.cat_vars)[0]
x_num = ord_to_num(x_ord, cf.d_abs)
# check gradients
grads = cf.get_gradients(x_num, y, num_shape[1:], cf.cat_vars_ord)
assert grads.shape == num_shape
Loading

0 comments on commit a39d198

Please sign in to comment.