Skip to content
New issue

Have a question about this project? Sign up for a free GitHub account to open an issue and contact its maintainers and the community.

By clicking “Sign up for GitHub”, you agree to our terms of service and privacy statement. We’ll occasionally send you account related emails.

Already on GitHub? Sign in to your account

ENH: Check input dimensions against the initialized model_ #143

Open
wants to merge 21 commits into
base: master
Choose a base branch
from
Open
Show file tree
Hide file tree
Changes from 13 commits
Commits
Show all changes
21 commits
Select commit Hold shift + click to select a range
File filter

Filter by extension

Filter by extension

Conversations
Failed to load comments.
Loading
Jump to
Jump to file
Failed to load files.
Loading
Diff view
Diff view
12 changes: 7 additions & 5 deletions scikeras/_utils.py
Original file line number Diff line number Diff line change
Expand Up @@ -4,7 +4,7 @@
import warnings

from inspect import isclass
from typing import Any, Callable, Dict, Iterable, List, Union
from typing import Any, Callable, Dict, Iterable, List, Mapping, Union

import numpy as np
import tensorflow as tf
Expand Down Expand Up @@ -107,16 +107,18 @@ def make_model_picklable(model_obj):


def _windows_upcast_ints(
arr: Union[List[np.ndarray], np.ndarray]
inp: Union[List[np.ndarray], Dict[Any, np.ndarray], np.ndarray]
) -> Union[List[np.ndarray], np.ndarray]:
# see tensorflow/probability#886
def _upcast(x):
return x.astype("int64") if x.dtype == np.int32 else x

if isinstance(arr, np.ndarray):
return _upcast(arr)
if isinstance(inp, np.ndarray):
return _upcast(inp)
elif isinstance(inp, Mapping):
return {k: _upcast(x_) for k, x_ in inp.items()}
else:
return [_upcast(x_) for x_ in arr]
return [_upcast(x_) for x_ in inp]


def route_params(
Expand Down
22 changes: 19 additions & 3 deletions scikeras/utils/transformers.py
Original file line number Diff line number Diff line change
Expand Up @@ -184,10 +184,10 @@ def fit(self, y: np.ndarray) -> "ClassifierLabelEncoder":
target_type = "multiclass-onehot"

self.n_outputs_ = 1
self.n_outputs_expected_ = 1
self._y_dtype = y.dtype
self._target_type = target_type

# Record class metadata
if target_type in ("binary", "multiclass"):
self.classes_ = self._final_encoder[1].categories_[0]
self.n_classes_ = self.classes_.size
Expand All @@ -198,6 +198,23 @@ def fit(self, y: np.ndarray) -> "ClassifierLabelEncoder":
self.classes_ = None
self.n_classes_ = None

# Record output metadata
if target_type == "binary":
# single output, single output unit, sigmoid activation
# it is also possible to use 2 units with a softmax activation,
# but there's generally no reason to do that so we assume a single unit
self.n_outputs_expected_ = 1
elif target_type in ("multiclass", "multiclass-onehot", "multilabel-indicator"):
# single output, one unit per class
self.n_outputs_expected_ = self.n_classes_
elif target_type == "multiclass-multioutput":
# This could be anything! We don't try to guess
# For user-written transformers, the expected format is
# List[int] (where the indexes are output numbers, as in Model.outputs)
# or
# Dict[str, int] (where the keys are output names, as in Model.output_names)
self.n_outputs_expected_ = None

return self

def transform(self, y: np.ndarray) -> np.ndarray:
Expand Down Expand Up @@ -328,8 +345,7 @@ def fit(self, y: np.ndarray) -> "RegressorTargetEncoder":
"""
self._y_dtype = y.dtype
self._y_shape = y.shape
self.n_outputs_ = 1 if y.ndim == 1 else y.shape[1]
self.n_outputs_expected_ = 1
self.n_outputs_ = self.n_outputs_expected_ = 1 if y.ndim == 1 else y.shape[1]
return self

def transform(self, y: np.ndarray) -> np.ndarray:
Expand Down
96 changes: 81 additions & 15 deletions scikeras/wrappers.py
Original file line number Diff line number Diff line change
Expand Up @@ -5,7 +5,7 @@
import warnings

from collections import defaultdict
from typing import Any, Callable, Dict, Iterable, List, Tuple, Type, Union
from typing import Any, Callable, Dict, Iterable, List, Mapping, Tuple, Type, Union

import numpy as np
import tensorflow as tf
Expand Down Expand Up @@ -530,22 +530,78 @@ def _fit_keras_model(
raise e
self.history_[key] += val

def _check_model_compatibility(self, y: np.ndarray) -> None:
"""Checks that the model output number and y shape match.
def _check_model_outputs(self, y):
# output shapes depend on the number of classes in classification,
# hence we cannot just check y here, we need the user (or the
# data transformer) to tell us what to expect via n_outputs_expected_

# n_outputs_expected_ is generated by data transformers
# and recovered via target_encoder_.get_meta();
# we recognize it but do not force it to be
# generated to avoid forcing users to subclass
# generic transformers just to make them SciKeras compatible

# n_outputs_expected_ can be of the following types:
# int (for a single output, this is the number of output units)
# List[int] (where the indexes are output numbers, as in Model.outputs)
# or
# Dict[str, int] (where the keys are output names, as in Model.output_names)
exp = getattr(self, "n_outputs_expected_", None)
# for output layers, the number of units is the second dimension
# of the shape (the first is samples / batch size)
actual = [output.shape[1] for output in self.model_.outputs]
actual_names = self.model_.output_names
if isinstance(exp, int):
# Keras outputs get saved as a list with a single item
# even when given as a bare object
exp = [exp]
if isinstance(exp, dict):
if set(exp.keys()) != set(actual_names):
raise ValueError(
"Output names in `n_outputs_expected_` do not"
"match the Model's output names:"
f"\n {set(exp.keys())}"
"\n vs"
f"\n {set(actual_names)}"
)
# Convert to a list to re-use validation code
exp = [exp[outname] for outname in actual_names]
if isinstance(exp, list):
if not len(exp) == len(actual):
raise ValueError(
"The target ``y`` seems to consist of"
f" {len(exp)} outputs, but this Keras"
f" Model has {len(actual)} outputs."
)
for n, a, e in zip(actual_names, actual, exp):
if a != e:
raise ValueError(
f"Model output {n} expected to have {e} output units"
f" but got {a}!"
)

This is in place to avoid cryptic TF errors.
"""
# check if this is a multi-output model
if hasattr(self, "n_outputs_expected_"):
# n_outputs_expected_ is generated by data transformers
# we recognize the attribute but do not force it to be
# generated
if self.n_outputs_expected_ != len(self.model_.outputs):
def _check_model_inputs(self, X):
if isinstance(X, np.ndarray):
# Keras Model's inputs are always a list
X = [X]
elif isinstance(X, Mapping):
X = [X[inp_name] for inp_name in self.model_.input_names]
if len(X) != len(self.model_.inputs):
raise ValueError(
f"``X`` has {len(X)} inputs, but the Keras model"
f" has {len(self.model_.inputs)} inputs."
)
for X_in, model_in in zip(X, self.model_.inputs):
# check shape
X_in_shape = (1,) if X_in.ndim == 1 else X_in.shape[1:]
model_in_shape = model_in.shape[1:]
if X_in_shape != model_in_shape:
raise ValueError(
"Detected a Keras model input of size"
f" {y[0].shape[0]}, but {self.model_} has"
f" {self.model_.outputs} outputs"
f"Input {model_in.name} expected shape"
f" {model_in_shape} but got {X_in_shape}."
)

def _check_model_loss(self):
# check that if the user gave us a loss function it ended up in
# the actual model
init_params = inspect.signature(self.__init__).parameters
Expand All @@ -565,6 +621,16 @@ def _check_model_compatibility(self, y: np.ndarray) -> None:
" Data may not match loss function!"
)

def _check_model_compatibility(self, X, y) -> None:
"""Checks that the model inputs, outputs and loss
match the given or expected X, y & loss.

This is in place to avoid cryptic TF errors.
"""
self._check_model_inputs(X)
self._check_model_outputs(y)
self._check_model_loss()

def _validate_data(
self, X=None, y=None, reset: bool = False
) -> Tuple[np.ndarray, Union[np.ndarray, None]]:
Expand Down Expand Up @@ -854,7 +920,7 @@ def _fit(
y = self.target_encoder_.transform(y)
X = self.feature_encoder_.transform(X)

self._check_model_compatibility(y)
self._check_model_compatibility(X, y)

self._fit_keras_model(
X,
Expand Down
7 changes: 2 additions & 5 deletions tests/mlp_models.py
Original file line number Diff line number Diff line change
Expand Up @@ -29,11 +29,8 @@ def dynamic_classifier(
out = [Dense(1, activation="sigmoid")(hidden)]
elif target_type_ == "multilabel-indicator":
compile_kwargs["loss"] = compile_kwargs["loss"] or "binary_crossentropy"
if isinstance(n_classes_, list):
out = [
Dense(1, activation="sigmoid")(hidden)
for _ in range(n_outputs_expected_)
]
if isinstance(n_outputs_expected_, list):
out = [Dense(n, activation="sigmoid")(hidden) for n in n_outputs_expected_]
else:
out = Dense(n_classes_, activation="softmax")(hidden)
elif target_type_ == "multiclass-multioutput":
Expand Down
3 changes: 2 additions & 1 deletion tests/multi_output_models.py
Original file line number Diff line number Diff line change
Expand Up @@ -15,7 +15,8 @@ def fit(self, y: np.ndarray) -> "MultiLabelTransformer":
return super().fit(y)
# y = array([1, 1, 1, 0], [0, 0, 1, 1])
# each col will be processed as multiple binary classifications
self.n_outputs_ = self.n_outputs_expected_ = y.shape[1]
self.n_outputs_ = y.shape[1]
self.n_outputs_expected_ = [1] * y.shape[1]
self._y_dtype = y.dtype
self.classes_ = [np.array([0, 1])] * y.shape[1]
self.n_classes_ = [2] * y.shape[1]
Expand Down
21 changes: 9 additions & 12 deletions tests/test_api.py
Original file line number Diff line number Diff line change
Expand Up @@ -47,11 +47,9 @@ def build_fn_clf(
model.add(keras.layers.Activation("relu"))
model.add(keras.layers.Dense(hidden_dim))
model.add(keras.layers.Activation("relu"))
model.add(keras.layers.Dense(n_classes_))
model.add(keras.layers.Activation("softmax"))
model.compile(
optimizer="sgd", loss="sparse_categorical_crossentropy", metrics=["accuracy"]
)
model.add(keras.layers.Dense(1))
model.add(keras.layers.Activation("sigmoid"))
model.compile(optimizer="sgd", loss="binary_crossentropy", metrics=["accuracy"])
return model


Expand Down Expand Up @@ -156,7 +154,8 @@ def build_fn_clss(
model.add(Dense(X_shape_[1], activation="relu", input_shape=X_shape_[1:]))
for size in hidden_layer_sizes:
model.add(Dense(size, activation="relu"))
model.add(Dense(1, activation="softmax"))
model.add(keras.layers.Dense(1))
model.add(keras.layers.Activation("sigmoid"))
model.compile("adam", loss="binary_crossentropy", metrics=["accuracy"])
return model

Expand Down Expand Up @@ -314,13 +313,12 @@ def test_basic(self, config):
data = loader()
x_train, y_train = data.data[:100], data.target[:100]

n_classes_ = np.unique(y_train).size
# make y the same shape as will be used by .fit
if config != "MLPRegressor":
y_train = to_categorical(y_train)
meta = {
"n_classes_": n_classes_,
"target_type_": "multiclass",
"n_classes_": 2,
"target_type_": "binary",
"n_features_in_": x_train.shape[1],
"n_outputs_expected_": 1,
}
Expand Down Expand Up @@ -350,13 +348,12 @@ def test_ensemble(self, config):
data = loader()
x_train, y_train = data.data[:100], data.target[:100]

n_classes_ = np.unique(y_train).size
# make y the same shape as will be used by .fit
if config != "MLPRegressor":
y_train = to_categorical(y_train)
meta = {
"n_classes_": n_classes_,
"target_type_": "multiclass",
"n_classes_": 2,
"target_type_": "binary",
"n_features_in_": x_train.shape[1],
"n_outputs_expected_": 1,
}
Expand Down
8 changes: 5 additions & 3 deletions tests/test_errors.py
Original file line number Diff line number Diff line change
Expand Up @@ -22,14 +22,16 @@ def test_X_shape_change():
loss=KerasRegressor.r_squared,
hidden_layer_sizes=(100,),
)
X = np.array([[1, 2], [3, 4]]).reshape(2, 2, 1)
X = np.array([[1, 2], [3, 4]])
y = np.array([[0, 1, 0], [1, 0, 0]])

estimator.fit(X=X, y=y)

# Calling with a different number of dimensions for X raises an error
with pytest.raises(ValueError, match="dimensions in X"):
# Calling with a different number of dimensions for X raises an error
estimator.partial_fit(X=X.reshape(2, 2), y=y)
estimator.partial_fit(X=X.reshape(2, 2, 1), y=y)
with pytest.raises(ValueError, match="dimensions in X"):
estimator.predict(X=X.reshape(2, 2, 1))


def test_unknown_param():
Expand Down
Loading