Skip to content

[MRG] New API should allow prediction functions and scoring #95

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

Merged
Show file tree
Hide file tree
Changes from 6 commits
Commits
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
111 changes: 109 additions & 2 deletions metric_learn/base_metric.py
Original file line number Diff line number Diff line change
@@ -1,9 +1,11 @@
from numpy.linalg import inv, cholesky
from numpy.linalg import cholesky
from sklearn.base import BaseEstimator, TransformerMixin
from sklearn.utils.validation import check_array
from sklearn.metrics import roc_auc_score
import numpy as np


class BaseMetricLearner(BaseEstimator, TransformerMixin):
class BaseMetricLearner(BaseEstimator):
def __init__(self):
raise NotImplementedError('BaseMetricLearner should not be instantiated')

Expand All @@ -30,6 +32,9 @@ def transformer(self):
"""
return cholesky(self.metric()).T


class MetricTransformer(TransformerMixin):

def transform(self, X=None):
"""Applies the metric transformation.

Expand All @@ -49,3 +54,105 @@ def transform(self, X=None):
X = check_array(X, accept_sparse=True)
L = self.transformer()
return X.dot(L.T)


class _PairsClassifierMixin:

def predict(self, pairs):
"""Predicts the learned similarity between input pairs.
Copy link
Member

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

should be metric instead of similarity here

Copy link
Member Author

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

Yes indeed, thanks


Returns the learned metric value between samples in every pair. It should
ideally be low for similar samples and high for dissimilar samples.

Parameters
----------
pairs : array-like, shape=(n_constraints, 2, n_features)
A constrained dataset of paired samples.

Returns
-------
y_predicted : `numpy.ndarray` of floats, shape=(n_constraints,)
The predicted learned metric value between samples in every pair.
"""
pairwise_diffs = pairs[:, 0, :] - pairs[:, 1, :]
return np.sqrt(np.sum(pairwise_diffs.dot(self.metric()) * pairwise_diffs,
axis=1))

def decision_function(self, pairs):
return self.predict(pairs)

def score(self, pairs, y):
"""Computes score of pairs similarity prediction.

Returns the ``roc_auc`` score of the fitted metric learner. It is
computed in the following way: for every value of a threshold
``t`` we classify all pairs of samples where the predicted distance is
inferior to ``t`` as belonging to the "similar" class, and the other as
belonging to the "dissimilar" class, and we count false positive and
true positives as in a classical ``roc_auc`` curve.

Parameters
----------
pairs : array-like, shape=(n_constraints, 2, n_features)
Input Pairs.

y : array-like, shape=(n_constraints,)
The corresponding labels.

Returns
-------
score : float
The ``roc_auc`` score.
"""
return roc_auc_score(y, self.decision_function(pairs))


class _QuadrupletsClassifierMixin:

def predict(self, quadruplets):
"""Predicts differences between sample similarities in input quadruplets.
Copy link
Member

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

distances?

Copy link
Member Author

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

Yes, thanks


For each quadruplet of samples, computes the difference between the learned
metric of the first pair minus the learned metric of the second pair.

Parameters
----------
quadruplets : array-like, shape=(n_constraints, 4, n_features)
Input quadruplets.

Returns
-------
prediction : np.ndarray of floats, shape=(n_constraints,)
Metric differences.
"""
similar_diffs = quadruplets[:, 0, :] - quadruplets[:, 1, :]
dissimilar_diffs = quadruplets[:, 2, :] - quadruplets[:, 3, :]
return (np.sqrt(np.sum(similar_diffs.dot(self.metric()) *
similar_diffs, axis=1)) -
np.sqrt(np.sum(dissimilar_diffs.dot(self.metric()) *
dissimilar_diffs, axis=1)))
Copy link
Contributor

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

This pattern, distance under some metric, seems like it should be factored out.

Copy link
Member Author

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

Yes indeed, the function will call function score_pairs (that returns the new metric between points) that will be inherited from the BaseMetricLearner, and implemented through ExplicitMixin (a Mixin for all learners that can embed data) (so score_pairs will be implemented as the euclidean distance between embeddings)

Copy link
Member

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

(this should ultimately be in the Mahalanobis Mixin)


def decision_function(self, quadruplets):
return self.predict(quadruplets)

def score(self, quadruplets, y=None):
"""Computes score on an input constrained dataset

Returns the accuracy score of the following classification task: a record
is correctly classified if the predicted similarity between the first two
samples is higher than that of the last two.

Parameters
----------
quadruplets : array-like, shape=(n_constraints, 4, n_features)
Input quadruplets.

y : Ignored, for scikit-learn compatibility.

Returns
-------
score : float
The quadruplets score.
"""
predicted_sign = self.decision_function(quadruplets) < 0
return np.sum(predicted_sign) / predicted_sign.shape[0]
Copy link
Contributor

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

Why not np.mean(np.sign(...)) here?

Copy link
Member Author

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

Much cleaner indeed, thanks !

4 changes: 2 additions & 2 deletions metric_learn/covariance.py
Original file line number Diff line number Diff line change
Expand Up @@ -12,10 +12,10 @@
import numpy as np
from sklearn.utils.validation import check_array

from .base_metric import BaseMetricLearner
from .base_metric import BaseMetricLearner, MetricTransformer


class Covariance(BaseMetricLearner):
class Covariance(BaseMetricLearner, MetricTransformer):
def __init__(self):
pass

Expand Down
56 changes: 31 additions & 25 deletions metric_learn/itml.py
Original file line number Diff line number Diff line change
Expand Up @@ -19,12 +19,13 @@
from sklearn.metrics import pairwise_distances
from sklearn.utils.validation import check_array, check_X_y

from .base_metric import BaseMetricLearner
from .base_metric import (BaseMetricLearner, _PairsClassifierMixin,
MetricTransformer)
from .constraints import Constraints, wrap_pairs
from ._util import vector_norm


class ITML(BaseMetricLearner):
class _BaseITML(BaseMetricLearner):
"""Information Theoretic Metric Learning (ITML)"""
def __init__(self, gamma=1., max_iter=1000, convergence_threshold=1e-3,
A0=None, verbose=False):
Expand Down Expand Up @@ -78,24 +79,7 @@ def _process_pairs(self, pairs, y, bounds):
y = np.hstack([np.ones(len(pos_pairs)), - np.ones(len(neg_pairs))])
return pairs, y


def fit(self, pairs, y, bounds=None):
"""Learn the ITML model.

Parameters
----------
pairs: array-like, shape=(n_constraints, 2, n_features)
Array of pairs. Each row corresponds to two points.
y: array-like, of shape (n_constraints,)
Labels of constraints. Should be -1 for dissimilar pair, 1 for similar.
bounds : list (pos,neg) pairs, optional
bounds on similarity, s.t. d(X[a],X[b]) < pos and d(X[c],X[d]) > neg

Returns
-------
self : object
Returns the instance.
"""
def _fit(self, pairs, y, bounds=None):
pairs, y = self._process_pairs(pairs, y, bounds)
gamma = self.gamma
pos_pairs, neg_pairs = pairs[y == 1], pairs[y == -1]
Expand Down Expand Up @@ -151,7 +135,29 @@ def metric(self):
return self.A_


class ITML_Supervised(ITML):
class ITML(_BaseITML, _PairsClassifierMixin):

def fit(self, pairs, y, bounds=None):
"""Learn the ITML model.

Parameters
----------
pairs: array-like, shape=(n_constraints, 2, n_features)
Array of pairs. Each row corresponds to two points.
y: array-like, of shape (n_constraints,)
Labels of constraints. Should be -1 for dissimilar pair, 1 for similar.
bounds : list (pos,neg) pairs, optional
bounds on similarity, s.t. d(X[a],X[b]) < pos and d(X[c],X[d]) > neg

Returns
-------
self : object
Returns the instance.
"""
return self._fit(pairs, y, bounds=bounds)


class ITML_Supervised(_BaseITML, MetricTransformer):
"""Information Theoretic Metric Learning (ITML)"""
def __init__(self, gamma=1., max_iter=1000, convergence_threshold=1e-3,
num_labeled=np.inf, num_constraints=None, bounds=None, A0=None,
Expand All @@ -175,9 +181,9 @@ def __init__(self, gamma=1., max_iter=1000, convergence_threshold=1e-3,
verbose : bool, optional
if True, prints information while learning
"""
ITML.__init__(self, gamma=gamma, max_iter=max_iter,
convergence_threshold=convergence_threshold,
A0=A0, verbose=verbose)
_BaseITML.__init__(self, gamma=gamma, max_iter=max_iter,
convergence_threshold=convergence_threshold,
A0=A0, verbose=verbose)
self.num_labeled = num_labeled
self.num_constraints = num_constraints
self.bounds = bounds
Expand Down Expand Up @@ -207,4 +213,4 @@ def fit(self, X, y, random_state=np.random):
pos_neg = c.positive_negative_pairs(num_constraints,
random_state=random_state)
pairs, y = wrap_pairs(X, pos_neg)
return ITML.fit(self, pairs, y, bounds=self.bounds)
return _BaseITML._fit(self, pairs, y, bounds=self.bounds)
4 changes: 2 additions & 2 deletions metric_learn/lfda.py
Original file line number Diff line number Diff line change
Expand Up @@ -18,10 +18,10 @@
from sklearn.metrics import pairwise_distances
from sklearn.utils.validation import check_X_y

from .base_metric import BaseMetricLearner
from .base_metric import BaseMetricLearner, MetricTransformer


class LFDA(BaseMetricLearner):
class LFDA(BaseMetricLearner, MetricTransformer):
'''
Local Fisher Discriminant Analysis for Supervised Dimensionality Reduction
Sugiyama, ICML 2006
Expand Down
4 changes: 2 additions & 2 deletions metric_learn/lmnn.py
Original file line number Diff line number Diff line change
Expand Up @@ -17,11 +17,11 @@
from sklearn.utils.validation import check_X_y, check_array
from sklearn.metrics import euclidean_distances

from .base_metric import BaseMetricLearner
from .base_metric import BaseMetricLearner, MetricTransformer


# commonality between LMNN implementations
class _base_LMNN(BaseMetricLearner):
class _base_LMNN(BaseMetricLearner, MetricTransformer):
def __init__(self, k=3, min_iter=50, max_iter=1000, learn_rate=1e-7,
regularization=0.5, convergence_tol=0.001, use_pca=True,
verbose=False):
Expand Down
60 changes: 34 additions & 26 deletions metric_learn/lsml.py
Original file line number Diff line number Diff line change
Expand Up @@ -13,11 +13,12 @@
from six.moves import xrange
from sklearn.utils.validation import check_array, check_X_y

from .base_metric import BaseMetricLearner
from .constraints import Constraints, wrap_pairs
from .base_metric import (BaseMetricLearner, _QuadrupletsClassifierMixin,
MetricTransformer)
from .constraints import Constraints


class LSML(BaseMetricLearner):
class _BaseLSML(BaseMetricLearner):
def __init__(self, tol=1e-3, max_iter=1000, prior=None, verbose=False):
"""Initialize LSML.

Expand Down Expand Up @@ -60,24 +61,7 @@ def _prepare_quadruplets(self, quadruplets, weights):
def metric(self):
return self.M_

def fit(self, quadruplets, weights=None):
"""Learn the LSML model.

Parameters
----------
quadruplets : array-like, shape=(n_constraints, 4, n_features)
Each row corresponds to 4 points. In order to supervise the
algorithm in the right way, we should have the four samples ordered
in a way such that: d(pairs[i, 0],X[i, 1]) < d(X[i, 2], X[i, 3])
for all 0 <= i < n_constraints.
weights : (n_constraints,) array of floats, optional
scale factor for each constraint

Returns
-------
self : object
Returns the instance.
"""
def _fit(self, quadruplets, weights=None):
self._prepare_quadruplets(quadruplets, weights)
step_sizes = np.logspace(-10, 0, 10)
# Keep track of the best step size and the loss at that step.
Expand Down Expand Up @@ -140,7 +124,30 @@ def _gradient(self, metric):
return dMetric


class LSML_Supervised(LSML):
class LSML(_BaseLSML, _QuadrupletsClassifierMixin):

def fit(self, quadruplets, weights=None):
"""Learn the LSML model.

Parameters
----------
quadruplets : array-like, shape=(n_constraints, 4, n_features)
Each row corresponds to 4 points. In order to supervise the
algorithm in the right way, we should have the four samples ordered
in a way such that: d(pairs[i, 0],X[i, 1]) < d(X[i, 2], X[i, 3])
for all 0 <= i < n_constraints.
weights : (n_constraints,) array of floats, optional
scale factor for each constraint

Returns
-------
self : object
Returns the instance.
"""
return self._fit(quadruplets, weights=weights)


class LSML_Supervised(_BaseLSML, MetricTransformer):
def __init__(self, tol=1e-3, max_iter=1000, prior=None, num_labeled=np.inf,
num_constraints=None, weights=None, verbose=False):
"""Initialize the learner.
Expand All @@ -160,8 +167,8 @@ def __init__(self, tol=1e-3, max_iter=1000, prior=None, num_labeled=np.inf,
verbose : bool, optional
if True, prints information while learning
"""
LSML.__init__(self, tol=tol, max_iter=max_iter, prior=prior,
verbose=verbose)
_BaseLSML.__init__(self, tol=tol, max_iter=max_iter, prior=prior,
verbose=verbose)
self.num_labeled = num_labeled
self.num_constraints = num_constraints
self.weights = weights
Expand Down Expand Up @@ -189,5 +196,6 @@ def fit(self, X, y, random_state=np.random):
c = Constraints.random_subset(y, self.num_labeled,
random_state=random_state)
pos_neg = c.positive_negative_pairs(num_constraints, same_length=True,
random_state=random_state)
return LSML.fit(self, X[np.column_stack(pos_neg)], weights=self.weights)
random_state=random_state)
return _BaseLSML._fit(self, X[np.column_stack(pos_neg)],
weights=self.weights)
4 changes: 2 additions & 2 deletions metric_learn/mlkr.py
Original file line number Diff line number Diff line change
Expand Up @@ -13,12 +13,12 @@
from sklearn.decomposition import PCA
from sklearn.utils.validation import check_X_y

from .base_metric import BaseMetricLearner
from .base_metric import BaseMetricLearner, MetricTransformer

EPS = np.finfo(float).eps


class MLKR(BaseMetricLearner):
class MLKR(BaseMetricLearner, MetricTransformer):
"""Metric Learning for Kernel Regression (MLKR)"""
def __init__(self, num_dims=None, A0=None, epsilon=0.01, alpha=0.0001,
max_iter=1000):
Expand Down
Loading