Skip to content

Commit

Permalink
MAINT unpack 0-dim NumPy array instead of implicit conversion (scikit…
Browse files Browse the repository at this point in the history
…-learn#26345)

Co-authored-by: Jérémie du Boisberranger <[email protected]>
  • Loading branch information
glemaitre and jeremiedbb committed Jun 9, 2023
1 parent 66733c4 commit 9eea5b7
Show file tree
Hide file tree
Showing 7 changed files with 18 additions and 10 deletions.
3 changes: 3 additions & 0 deletions sklearn/linear_model/_logistic.py
Original file line number Diff line number Diff line change
Expand Up @@ -506,6 +506,9 @@ def _logistic_regression_path(
w0 = np.concatenate([coef_.ravel(), intercept_])
else:
w0 = coef_.ravel()
# n_iter_i is an array for each class. However, `target` is always encoded
# in {-1, 1}, so we only take the first element of n_iter_i.
n_iter_i = n_iter_i.item()

elif solver in ["sag", "saga"]:
if multi_class == "multinomial":
Expand Down
4 changes: 2 additions & 2 deletions sklearn/linear_model/tests/test_sag.py
Original file line number Diff line number Diff line change
Expand Up @@ -95,7 +95,7 @@ def sag(

for epoch in range(n_iter):
for k in range(n_samples):
idx = int(rng.rand(1) * n_samples)
idx = int(rng.rand() * n_samples)
# idx = k
entry = X[idx]
seen.add(idx)
Expand Down Expand Up @@ -167,7 +167,7 @@ def sag_sparse(
for epoch in range(n_iter):
for k in range(n_samples):
# idx = k
idx = int(rng.rand(1) * n_samples)
idx = int(rng.rand() * n_samples)
entry = X[idx]
seen.add(idx)

Expand Down
3 changes: 2 additions & 1 deletion sklearn/metrics/tests/test_pairwise.py
Original file line number Diff line number Diff line change
Expand Up @@ -298,7 +298,8 @@ def test_pairwise_precomputed_non_negative():
def callable_rbf_kernel(x, y, **kwds):
# Callable version of pairwise.rbf_kernel.
K = rbf_kernel(np.atleast_2d(x), np.atleast_2d(y), **kwds)
return K
# unpack the output since this is a scalar packed in a 0-dim array
return K.item()


@pytest.mark.parametrize(
Expand Down
2 changes: 1 addition & 1 deletion sklearn/mixture/tests/test_bayesian_mixture.py
Original file line number Diff line number Diff line change
Expand Up @@ -58,7 +58,7 @@ def test_log_wishart_norm():
),
0,
)
)
).item()
predected_norm = _log_wishart_norm(
degrees_of_freedom, log_det_precisions_chol, n_features
)
Expand Down
2 changes: 1 addition & 1 deletion sklearn/tests/test_discriminant_analysis.py
Original file line number Diff line number Diff line change
Expand Up @@ -187,7 +187,7 @@ def generate_dataset(n_samples, centers, covariances, random_state=None):
sample = np.array([[-22, 22]])

def discriminant_func(sample, coef, intercept, clazz):
return np.exp(intercept[clazz] + np.dot(sample, coef[clazz]))
return np.exp(intercept[clazz] + np.dot(sample, coef[clazz])).item()

prob = np.array(
[
Expand Down
11 changes: 7 additions & 4 deletions sklearn/tree/_export.py
Original file line number Diff line number Diff line change
Expand Up @@ -11,6 +11,7 @@
# Li Li <[email protected]>
# Giuseppe Vettigli <[email protected]>
# License: BSD 3 clause
from collections.abc import Iterable
from io import StringIO
from numbers import Integral

Expand Down Expand Up @@ -247,7 +248,7 @@ def get_color(self, value):
color = list(self.colors["rgb"][np.argmax(value)])
sorted_values = sorted(value, reverse=True)
if len(sorted_values) == 1:
alpha = 0
alpha = 0.0
else:
alpha = (sorted_values[0] - sorted_values[1]) / (1 - sorted_values[1])
else:
Expand All @@ -256,8 +257,6 @@ def get_color(self, value):
alpha = (value - self.colors["bounds"][0]) / (
self.colors["bounds"][1] - self.colors["bounds"][0]
)
# unpack numpy scalars
alpha = float(alpha)
# compute the color as alpha against white
color = [int(round(alpha * c + (1 - alpha) * 255, 0)) for c in color]
# Return html color code in #RRGGBB format
Expand All @@ -277,8 +276,12 @@ def get_fill_color(self, tree, node_id):
if tree.n_outputs == 1:
node_val = tree.value[node_id][0, :] / tree.weighted_n_node_samples[node_id]
if tree.n_classes[0] == 1:
# Regression
# Regression or degraded classification with single class
node_val = tree.value[node_id][0, :]
if isinstance(node_val, Iterable) and self.colors["bounds"] is not None:
# Only unpack the float only for the regression tree case.
# Classification tree requires an Iterable in `get_color`.
node_val = node_val.item()
else:
# If multi-output color node by impurity
node_val = -tree.impurity[node_id]
Expand Down
3 changes: 2 additions & 1 deletion sklearn/utils/random.py
Original file line number Diff line number Diff line change
Expand Up @@ -75,7 +75,8 @@ def _random_choice_csc(n_samples, classes, class_probability=None, random_state=
# If there are nonzero classes choose randomly using class_probability
rng = check_random_state(random_state)
if classes[j].shape[0] > 1:
p_nonzero = 1 - class_prob_j[classes[j] == 0]
index_class_0 = np.flatnonzero(classes[j] == 0).item()
p_nonzero = 1 - class_prob_j[index_class_0]
nnz = int(n_samples * p_nonzero)
ind_sample = sample_without_replacement(
n_population=n_samples, n_samples=nnz, random_state=random_state
Expand Down

0 comments on commit 9eea5b7

Please sign in to comment.