Skip to content
Merged
Show file tree
Hide file tree
Changes from all 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
25 changes: 18 additions & 7 deletions ignite/contrib/metrics/regression/manhattan_distance.py
Original file line number Diff line number Diff line change
@@ -1,32 +1,43 @@
from typing import Callable, Union

import torch

from ignite.contrib.metrics.regression._base import _BaseRegression
from ignite.metrics.metric import reinit__is_reduced, sync_all_reduce


class ManhattanDistance(_BaseRegression):
r"""
Calculates the Manhattan Distance:

:math:`\text{MD} = \sum_{j=1}^n (A_j - P_j)`,
:math:`\text{MD} = \sum_{j=1}^n |A_j - P_j|`,

where :math:`A_j` is the ground truth and :math:`P_j` is the predicted value.

More details can be found in `Botchkarev 2018`__.
More details can be found in `scikit-learn distance metrics`__.

- ``update`` must receive output of the form ``(y_pred, y)`` or ``{'y_pred': y_pred, 'y': y}``.
- `y` and `y_pred` must be of same shape `(N, )` or `(N, 1)`.

__ https://arxiv.org/abs/1809.03006
__ https://scikit-learn.org/stable/modules/generated/sklearn.neighbors.DistanceMetric.html

"""

def __init__(
self, output_transform: Callable = lambda x: x, device: Union[str, torch.device] = torch.device("cpu")
):
self._sum_of_errors = None
super(ManhattanDistance, self).__init__(output_transform, device)

@reinit__is_reduced
def reset(self):
self._sum_of_errors = 0.0
self._sum_of_errors = torch.tensor(0.0, device=self._device)

def _update(self, output):
y_pred, y = output
errors = y.view_as(y_pred) - y_pred
self._sum_of_errors += torch.sum(errors).item()
errors = torch.abs(y - y_pred)
self._sum_of_errors += torch.sum(errors).to(self._device)

@sync_all_reduce("_sum_of_errors")
def compute(self):
return self._sum_of_errors
return self._sum_of_errors.item()
126 changes: 118 additions & 8 deletions tests/ignite/contrib/metrics/regression/test_manhattan_distance.py
Original file line number Diff line number Diff line change
@@ -1,7 +1,11 @@
import os

import numpy as np
import pytest
import torch
from sklearn.neighbors import DistanceMetric

import ignite.distributed as idist
from ignite.contrib.metrics.regression import ManhattanDistance


Expand Down Expand Up @@ -30,18 +34,124 @@ def test_mahattan_distance():

m = ManhattanDistance()

manhattan = DistanceMetric.get_metric("manhattan")

m.update((torch.from_numpy(a), torch.from_numpy(ground_truth)))
np_ans = (ground_truth - a).sum()
assert m.compute() == pytest.approx(np_ans)
np_sum = np.abs(ground_truth - a).sum()
assert m.compute() == pytest.approx(np_sum)
assert manhattan.pairwise([a, ground_truth])[0][1] == pytest.approx(np_sum)

m.update((torch.from_numpy(b), torch.from_numpy(ground_truth)))
np_ans += (ground_truth - b).sum()
assert m.compute() == pytest.approx(np_ans)
np_sum += np.abs(ground_truth - b).sum()
assert m.compute() == pytest.approx(np_sum)
v1 = np.hstack([a, b])
v2 = np.hstack([ground_truth, ground_truth])
assert manhattan.pairwise([v1, v2])[0][1] == pytest.approx(np_sum)

m.update((torch.from_numpy(c), torch.from_numpy(ground_truth)))
np_ans += (ground_truth - c).sum()
assert m.compute() == pytest.approx(np_ans)
np_sum += np.abs(ground_truth - c).sum()
assert m.compute() == pytest.approx(np_sum)
v1 = np.hstack([v1, c])
v2 = np.hstack([v2, ground_truth])
assert manhattan.pairwise([v1, v2])[0][1] == pytest.approx(np_sum)

m.update((torch.from_numpy(d), torch.from_numpy(ground_truth)))
np_ans += (ground_truth - d).sum()
assert m.compute() == pytest.approx(np_ans)
np_sum += np.abs(ground_truth - d).sum()
assert m.compute() == pytest.approx(np_sum)
v1 = np.hstack([v1, d])
v2 = np.hstack([v2, ground_truth])
assert manhattan.pairwise([v1, v2])[0][1] == pytest.approx(np_sum)


def _test_distrib_compute(device):
rank = idist.get_rank()

manhattan = DistanceMetric.get_metric("manhattan")

def _test(metric_device):
metric_device = torch.device(metric_device)
m = ManhattanDistance(device=metric_device)
torch.manual_seed(10 + rank)

y_pred = torch.randint(0, 10, size=(10,), device=device).float()
y = torch.randint(0, 10, size=(10,), device=device).float()

m.update((y_pred, y))

# gather y_pred, y
y_pred = idist.all_gather(y_pred)
y = idist.all_gather(y)

np_y_pred = y_pred.cpu().numpy()
np_y = y.cpu().numpy()
res = m.compute()
assert manhattan.pairwise([np_y_pred, np_y])[0][1] == pytest.approx(res)

for _ in range(3):
_test("cpu")
if device.type != "xla":
_test(idist.device())


@pytest.mark.distributed
@pytest.mark.skipif(not idist.has_native_dist_support, reason="Skip if no native dist support")
@pytest.mark.skipif(torch.cuda.device_count() < 1, reason="Skip if no GPU")
def test_distrib_gpu(distributed_context_single_node_nccl):
device = torch.device("cuda:{}".format(distributed_context_single_node_nccl["local_rank"]))
_test_distrib_compute(device)


@pytest.mark.distributed
@pytest.mark.skipif(not idist.has_native_dist_support, reason="Skip if no native dist support")
def test_distrib_cpu(distributed_context_single_node_gloo):

device = torch.device("cpu")
_test_distrib_compute(device)


@pytest.mark.distributed
@pytest.mark.skipif(not idist.has_hvd_support, reason="Skip if no Horovod dist support")
@pytest.mark.skipif("WORLD_SIZE" in os.environ, reason="Skip if launched as multiproc")
def test_distrib_hvd(gloo_hvd_executor):

device = torch.device("cpu" if not torch.cuda.is_available() else "cuda")
nproc = 4 if not torch.cuda.is_available() else torch.cuda.device_count()

gloo_hvd_executor(_test_distrib_compute, (device,), np=nproc, do_init=True)


@pytest.mark.multinode_distributed
@pytest.mark.skipif(not idist.has_native_dist_support, reason="Skip if no native dist support")
@pytest.mark.skipif("MULTINODE_DISTRIB" not in os.environ, reason="Skip if not multi-node distributed")
def test_multinode_distrib_cpu(distributed_context_multi_node_gloo):
device = torch.device("cpu")
_test_distrib_compute(device)


@pytest.mark.multinode_distributed
@pytest.mark.skipif(not idist.has_native_dist_support, reason="Skip if no native dist support")
@pytest.mark.skipif("GPU_MULTINODE_DISTRIB" not in os.environ, reason="Skip if not multi-node distributed")
def test_multinode_distrib_gpu(distributed_context_multi_node_nccl):
device = torch.device("cuda:{}".format(distributed_context_multi_node_nccl["local_rank"]))
_test_distrib_compute(device)


@pytest.mark.tpu
@pytest.mark.skipif("NUM_TPU_WORKERS" in os.environ, reason="Skip if NUM_TPU_WORKERS is in env vars")
@pytest.mark.skipif(not idist.has_xla_support, reason="Skip if no PyTorch XLA package")
def test_distrib_single_device_xla():
device = idist.device()
_test_distrib_compute(device)


def _test_distrib_xla_nprocs(index):
device = idist.device()
_test_distrib_compute(device)


@pytest.mark.tpu
@pytest.mark.skipif("NUM_TPU_WORKERS" not in os.environ, reason="Skip if no NUM_TPU_WORKERS in env vars")
@pytest.mark.skipif(not idist.has_xla_support, reason="Skip if no PyTorch XLA package")
def test_distrib_xla_nprocs(xmp_executor):
n = int(os.environ["NUM_TPU_WORKERS"])
xmp_executor(_test_distrib_xla_nprocs, args=(), nprocs=n)