Skip to content

Commit a0d77a5

Browse files
authored
Merge branch 'master' into iou-compute-in-cuda
2 parents bfe71c5 + f3caa76 commit a0d77a5

File tree

7 files changed

+110
-90
lines changed

7 files changed

+110
-90
lines changed

CHANGELOG.md

+3
Original file line numberDiff line numberDiff line change
@@ -23,6 +23,9 @@ and this project adheres to [Semantic Versioning](https://semver.org/spec/v2.0.0
2323
- Added `top_k` argument to `RetrievalMRR` in retrieval package ([#1961](https://github.com/Lightning-AI/torchmetrics/pull/1961))
2424

2525

26+
- Added warning to `MeanAveragePrecision` if too many detections are observed ([#1978](https://github.com/Lightning-AI/torchmetrics/pull/1978))
27+
28+
2629
### Changed
2730

2831
-

src/torchmetrics/detection/mean_ap.py

+20-2
Original file line numberDiff line numberDiff line change
@@ -24,6 +24,7 @@
2424

2525
from torchmetrics.detection.helpers import _fix_empty_tensors, _input_validator
2626
from torchmetrics.metric import Metric
27+
from torchmetrics.utilities import rank_zero_warn
2728
from torchmetrics.utilities.imports import (
2829
_MATPLOTLIB_AVAILABLE,
2930
_PYCOCOTOOLS_AVAILABLE,
@@ -239,6 +240,8 @@ class MeanAveragePrecision(Metric):
239240
groundtruth_crowds: List[Tensor]
240241
groundtruth_area: List[Tensor]
241242

243+
warn_on_many_detections: bool = True
244+
242245
def __init__(
243246
self,
244247
box_format: Literal["xyxy", "xywh", "cxcywh"] = "xyxy",
@@ -329,7 +332,7 @@ def update(self, preds: List[Dict[str, Tensor]], target: List[Dict[str, Tensor]]
329332
_input_validator(preds, target, iou_type=self.iou_type)
330333

331334
for item in preds:
332-
detections = self._get_safe_item_values(item)
335+
detections = self._get_safe_item_values(item, warn=self.warn_on_many_detections)
333336

334337
self.detections.append(detections)
335338
self.detection_labels.append(item["labels"])
@@ -542,11 +545,12 @@ def tm_to_coco(self, name: str = "tm_map_input") -> None:
542545
with open(f"{name}_target.json", "w") as f:
543546
f.write(target_json)
544547

545-
def _get_safe_item_values(self, item: Dict[str, Any]) -> Union[Tensor, Tuple]:
548+
def _get_safe_item_values(self, item: Dict[str, Any], warn: bool = False) -> Union[Tensor, Tuple]:
546549
"""Convert and return the boxes or masks from the item depending on the iou_type.
547550
548551
Args:
549552
item: input dictionary containing the boxes or masks
553+
warn: whether to warn if the number of boxes or masks exceeds the max_detection_thresholds
550554
551555
Returns:
552556
boxes or masks depending on the iou_type
@@ -556,12 +560,16 @@ def _get_safe_item_values(self, item: Dict[str, Any]) -> Union[Tensor, Tuple]:
556560
boxes = _fix_empty_tensors(item["boxes"])
557561
if boxes.numel() > 0:
558562
boxes = box_convert(boxes, in_fmt=self.box_format, out_fmt="xywh")
563+
if warn and len(boxes) > self.max_detection_thresholds[-1]:
564+
_warning_on_too_many_detections(self.max_detection_thresholds[-1])
559565
return boxes
560566
if self.iou_type == "segm":
561567
masks = []
562568
for i in item["masks"].cpu().numpy():
563569
rle = mask_utils.encode(np.asfortranarray(i))
564570
masks.append((tuple(rle["size"]), rle["counts"]))
571+
if warn and len(masks) > self.max_detection_thresholds[-1]:
572+
_warning_on_too_many_detections(self.max_detection_thresholds[-1])
565573
return tuple(masks)
566574
raise Exception(f"IOU type {self.iou_type} is not supported")
567575

@@ -747,3 +755,13 @@ def _gather_tuple_list(list_to_gather: List[Tuple], process_group: Optional[Any]
747755
dist.all_gather_object(list_gathered, list_to_gather, group=process_group)
748756

749757
return [list_gathered[rank][idx] for idx in range(len(list_gathered[0])) for rank in range(world_size)]
758+
759+
760+
def _warning_on_too_many_detections(limit: int) -> None:
761+
rank_zero_warn(
762+
f"Encountered more than {limit} detections in a single image. This means that certain detections with the"
763+
" lowest scores will be ignored, that may have an undesirable impact on performance. Please consider adjusting"
764+
" the `max_detection_threshold` to suit your use case. To disable this warning, set attribute class"
765+
" `warn_on_many_detections=False`, after initializing the metric.",
766+
UserWarning,
767+
)

src/torchmetrics/nominal/cramers.py

+15-21
Original file line numberDiff line numberDiff line change
@@ -39,11 +39,20 @@ class CramersV(Metric):
3939
\chi^2 = \sum_{i,j} \ frac{\left(n_{ij} - \frac{n_{i.} n_{.j}}{n}\right)^2}{\frac{n_{i.} n_{.j}}{n}}
4040
4141
where :math:`n_{ij}` denotes the number of times the values :math:`(A_i, B_j)` are observed with :math:`A_i, B_j`
42-
represent frequencies of values in ``preds`` and ``target``, respectively.
42+
represent frequencies of values in ``preds`` and ``target``, respectively. Cramer's V is a symmetric coefficient,
43+
i.e. :math:`V(preds, target) = V(target, preds)`, so order of input arguments does not matter. The output values
44+
lies in [0, 1] with 1 meaning the perfect association.
4345
44-
Cramer's V is a symmetric coefficient, i.e. :math:`V(preds, target) = V(target, preds)`.
46+
As input to ``forward`` and ``update`` the metric accepts the following input:
4547
46-
The output values lies in [0, 1] with 1 meaning the perfect association.
48+
- ``preds`` (:class:`~torch.Tensor`): Either 1D or 2D tensor of categorical (nominal) data from the first data
49+
series with shape ``(batch_size,)`` or ``(batch_size, num_classes)``, respectively.
50+
- ``target`` (:class:`~torch.Tensor`): Either 1D or 2D tensor of categorical (nominal) data from the second data
51+
series with shape ``(batch_size,)`` or ``(batch_size, num_classes)``, respectively.
52+
53+
As output of ``forward`` and ``compute`` the metric returns the following output:
54+
55+
- ``cramers_v`` (:class:`~torch.Tensor`): Scalar tensor containing the Cramer's V statistic.
4756
4857
Args:
4958
num_classes: Integer specifing the number of classes
@@ -52,16 +61,14 @@ class CramersV(Metric):
5261
nan_replace_value: Value to replace ``NaN``s when ``nan_strategy = 'replace'``
5362
kwargs: Additional keyword arguments, see :ref:`Metric kwargs` for more info.
5463
55-
Returns:
56-
Cramer's V statistic
57-
5864
Raises:
5965
ValueError:
6066
If `nan_strategy` is not one of `'replace'` and `'drop'`
6167
ValueError:
6268
If `nan_strategy` is equal to `'replace'` and `nan_replace_value` is not an `int` or `float`
6369
64-
Example:
70+
Example::
71+
6572
>>> from torchmetrics.nominal import CramersV
6673
>>> _ = torch.manual_seed(42)
6774
>>> preds = torch.randint(0, 4, (100,))
@@ -98,20 +105,7 @@ def __init__(
98105
self.add_state("confmat", torch.zeros(num_classes, num_classes), dist_reduce_fx="sum")
99106

100107
def update(self, preds: Tensor, target: Tensor) -> None:
101-
"""Update state with predictions and targets.
102-
103-
Args:
104-
preds: 1D or 2D tensor of categorical (nominal) data
105-
106-
- 1D shape: (batch_size,)
107-
- 2D shape: (batch_size, num_classes)
108-
109-
target: 1D or 2D tensor of categorical (nominal) data
110-
111-
- 1D shape: (batch_size,)
112-
- 2D shape: (batch_size, num_classes)
113-
114-
"""
108+
"""Update state with predictions and targets."""
115109
confmat = _cramers_v_update(preds, target, self.num_classes, self.nan_strategy, self.nan_replace_value)
116110
self.confmat += confmat
117111

src/torchmetrics/nominal/pearson.py

+16-23
Original file line numberDiff line numberDiff line change
@@ -43,30 +43,36 @@ class PearsonsContingencyCoefficient(Metric):
4343
.. math::
4444
\chi^2 = \sum_{i,j} \ frac{\left(n_{ij} - \frac{n_{i.} n_{.j}}{n}\right)^2}{\frac{n_{i.} n_{.j}}{n}}
4545
46-
where :math:`n_{ij}` denotes the number of times the values :math:`(A_i, B_j)` are observed
47-
with :math:`A_i, B_j` represent frequencies of values in ``preds`` and ``target``, respectively.
46+
where :math:`n_{ij}` denotes the number of times the values :math:`(A_i, B_j)` are observed with :math:`A_i, B_j`
47+
represent frequencies of values in ``preds`` and ``target``, respectively. Pearson's Contingency Coefficient is a
48+
symmetric coefficient, i.e. :math:`Pearson(preds, target) = Pearson(target, preds)`, so order of input arguments
49+
does not matter. The output values lies in [0, 1] with 1 meaning the perfect association.
4850
49-
Pearson's Contingency Coefficient is a symmetric coefficient, i.e.
50-
:math:`Pearson(preds, target) = Pearson(target, preds)`.
51+
As input to ``forward`` and ``update`` the metric accepts the following input:
5152
52-
The output values lies in [0, 1] with 1 meaning the perfect association.
53+
- ``preds`` (:class:`~torch.Tensor`): Either 1D or 2D tensor of categorical (nominal) data from the first data
54+
series with shape ``(batch_size,)`` or ``(batch_size, num_classes)``, respectively.
55+
- ``target`` (:class:`~torch.Tensor`): Either 1D or 2D tensor of categorical (nominal) data from the second data
56+
series with shape ``(batch_size,)`` or ``(batch_size, num_classes)``, respectively.
57+
58+
As output of ``forward`` and ``compute`` the metric returns the following output:
59+
60+
- ``pearsons_cc`` (:class:`~torch.Tensor`): Scalar tensor containing the Pearsons Contingency Coefficient statistic.
5361
5462
Args:
5563
num_classes: Integer specifing the number of classes
5664
nan_strategy: Indication of whether to replace or drop ``NaN`` values
5765
nan_replace_value: Value to replace ``NaN``s when ``nan_strategy = 'replace'``
5866
kwargs: Additional keyword arguments, see :ref:`Metric kwargs` for more info.
5967
60-
Returns:
61-
Pearson's Contingency Coefficient statistic
62-
6368
Raises:
6469
ValueError:
6570
If `nan_strategy` is not one of `'replace'` and `'drop'`
6671
ValueError:
6772
If `nan_strategy` is equal to `'replace'` and `nan_replace_value` is not an `int` or `float`
6873
69-
Example:
74+
Example::
75+
7076
>>> from torchmetrics.nominal import PearsonsContingencyCoefficient
7177
>>> _ = torch.manual_seed(42)
7278
>>> preds = torch.randint(0, 4, (100,))
@@ -101,20 +107,7 @@ def __init__(
101107
self.add_state("confmat", torch.zeros(num_classes, num_classes), dist_reduce_fx="sum")
102108

103109
def update(self, preds: Tensor, target: Tensor) -> None:
104-
"""Update state with predictions and targets.
105-
106-
Args:
107-
preds: 1D or 2D tensor of categorical (nominal) data:
108-
109-
- 1D shape: (batch_size,)
110-
- 2D shape: (batch_size, num_classes)
111-
112-
target: 1D or 2D tensor of categorical (nominal) data:
113-
114-
- 1D shape: (batch_size,)
115-
- 2D shape: (batch_size, num_classes)
116-
117-
"""
110+
"""Update state with predictions and targets."""
118111
confmat = _pearsons_contingency_coefficient_update(
119112
preds, target, self.num_classes, self.nan_strategy, self.nan_replace_value
120113
)

src/torchmetrics/nominal/theils_u.py

+16-18
Original file line numberDiff line numberDiff line change
@@ -34,23 +34,31 @@ class TheilsU(Metric):
3434
U(X|Y) = \frac{H(X) - H(X|Y)}{H(X)}
3535
3636
where :math:`H(X)` is entropy of variable :math:`X` while :math:`H(X|Y)` is the conditional entropy of :math:`X`
37-
given :math:`Y`. It is also know as the Uncertainty Coefficient.
37+
given :math:`Y`. It is also know as the Uncertainty Coefficient. Theils's U is an asymmetric coefficient, i.e.
38+
:math:`TheilsU(preds, target) \neq TheilsU(target, preds)`, so the order of the inputs matters. The output values
39+
lies in [0, 1], where a 0 means y has no information about x while value 1 means y has complete information about x.
3840
39-
Theils's U is an asymmetric coefficient, i.e. :math:`TheilsU(preds, target) \neq TheilsU(target, preds)`.
41+
As input to ``forward`` and ``update`` the metric accepts the following input:
4042
41-
The output values lies in [0, 1]. 0 means y has no information about x while value 1 means y has complete
42-
information about x.
43+
- ``preds`` (:class:`~torch.Tensor`): Either 1D or 2D tensor of categorical (nominal) data from the first data
44+
series (called X in the above definition) with shape ``(batch_size,)`` or ``(batch_size, num_classes)``,
45+
respectively.
46+
- ``target`` (:class:`~torch.Tensor`): Either 1D or 2D tensor of categorical (nominal) data from the second data
47+
series (called Y in the above definition) with shape ``(batch_size,)`` or ``(batch_size, num_classes)``,
48+
respectively.
49+
50+
As output of ``forward`` and ``compute`` the metric returns the following output:
51+
52+
- ``theils_u`` (:class:`~torch.Tensor`): Scalar tensor containing the Theil's U statistic.
4353
4454
Args:
4555
num_classes: Integer specifing the number of classes
4656
nan_strategy: Indication of whether to replace or drop ``NaN`` values
4757
nan_replace_value: Value to replace ``NaN``s when ``nan_strategy = 'replace'``
4858
kwargs: Additional keyword arguments, see :ref:`Metric kwargs` for more info.
4959
50-
Returns:
51-
Theil's U Statistic: Tensor
60+
Example::
5261
53-
Example:
5462
>>> from torchmetrics.nominal import TheilsU
5563
>>> _ = torch.manual_seed(42)
5664
>>> preds = torch.randint(10, (10,))
@@ -85,17 +93,7 @@ def __init__(
8593
self.add_state("confmat", torch.zeros(num_classes, num_classes), dist_reduce_fx="sum")
8694

8795
def update(self, preds: Tensor, target: Tensor) -> None:
88-
"""Update state with predictions and targets.
89-
90-
Args:
91-
preds: 1D or 2D tensor of categorical (nominal) data
92-
- 1D shape: (batch_size,)
93-
- 2D shape: (batch_size, num_classes)
94-
target: 1D or 2D tensor of categorical (nominal) data
95-
- 1D shape: (batch_size,)
96-
- 2D shape: (batch_size, num_classes)
97-
98-
"""
96+
"""Update state with predictions and targets."""
9997
confmat = _theils_u_update(preds, target, self.num_classes, self.nan_strategy, self.nan_replace_value)
10098
self.confmat += confmat
10199

src/torchmetrics/nominal/tschuprows.py

+16-22
Original file line numberDiff line numberDiff line change
@@ -38,12 +38,21 @@ class TschuprowsT(Metric):
3838
.. math::
3939
\chi^2 = \sum_{i,j} \ frac{\left(n_{ij} - \frac{n_{i.} n_{.j}}{n}\right)^2}{\frac{n_{i.} n_{.j}}{n}}
4040
41-
where :math:`n_{ij}` denotes the number of times the values :math:`(A_i, B_j)` are observed
42-
with :math:`A_i, B_j` represent frequencies of values in ``preds`` and ``target``, respectively.
41+
where :math:`n_{ij}` denotes the number of times the values :math:`(A_i, B_j)` are observed with :math:`A_i, B_j`
42+
represent frequencies of values in ``preds`` and ``target``, respectively. Tschuprow's T is a symmetric coefficient,
43+
i.e. :math:`T(preds, target) = T(target, preds)`, so order of input arguments does not matter. The output values
44+
lies in [0, 1] with 1 meaning the perfect association.
4345
44-
Tschuprow's T is a symmetric coefficient, i.e. :math:`T(preds, target) = T(target, preds)`.
46+
As input to ``forward`` and ``update`` the metric accepts the following input:
4547
46-
The output values lies in [0, 1] with 1 meaning the perfect association.
48+
- ``preds`` (:class:`~torch.Tensor`): Either 1D or 2D tensor of categorical (nominal) data from the first data
49+
series with shape ``(batch_size,)`` or ``(batch_size, num_classes)``, respectively.
50+
- ``target`` (:class:`~torch.Tensor`): Either 1D or 2D tensor of categorical (nominal) data from the second data
51+
series with shape ``(batch_size,)`` or ``(batch_size, num_classes)``, respectively.
52+
53+
As output of ``forward`` and ``compute`` the metric returns the following output:
54+
55+
- ``tschuprows_t`` (:class:`~torch.Tensor`): Scalar tensor containing the Tschuprow's T statistic.
4756
4857
Args:
4958
num_classes: Integer specifing the number of classes
@@ -52,16 +61,14 @@ class TschuprowsT(Metric):
5261
nan_replace_value: Value to replace ``NaN``s when ``nan_strategy = 'replace'``
5362
kwargs: Additional keyword arguments, see :ref:`Metric kwargs` for more info.
5463
55-
Returns:
56-
Tschuprow's T statistic
57-
5864
Raises:
5965
ValueError:
6066
If `nan_strategy` is not one of `'replace'` and `'drop'`
6167
ValueError:
6268
If `nan_strategy` is equal to `'replace'` and `nan_replace_value` is not an `int` or `float`
6369
64-
Example:
70+
Example::
71+
6572
>>> from torchmetrics.nominal import TschuprowsT
6673
>>> _ = torch.manual_seed(42)
6774
>>> preds = torch.randint(0, 4, (100,))
@@ -98,20 +105,7 @@ def __init__(
98105
self.add_state("confmat", torch.zeros(num_classes, num_classes), dist_reduce_fx="sum")
99106

100107
def update(self, preds: Tensor, target: Tensor) -> None:
101-
"""Update state with predictions and targets.
102-
103-
Args:
104-
preds: 1D or 2D tensor of categorical (nominal) data:
105-
106-
- 1D shape: (batch_size,)
107-
- 2D shape: (batch_size, num_classes)
108-
109-
target: 1D or 2D tensor of categorical (nominal) data:
110-
111-
- 1D shape: (batch_size,)
112-
- 2D shape: (batch_size, num_classes)
113-
114-
"""
108+
"""Update state with predictions and targets."""
115109
confmat = _tschuprows_t_update(preds, target, self.num_classes, self.nan_strategy, self.nan_replace_value)
116110
self.confmat += confmat
117111

tests/unittests/detection/test_map.py

+24-4
Original file line numberDiff line numberDiff line change
@@ -628,19 +628,19 @@ def test_error_on_wrong_input():
628628
)
629629

630630

631-
def _generate_random_segm_input(device):
631+
def _generate_random_segm_input(device, batch_size=2, num_preds_size=10, num_gt_size=10, random_size=True):
632632
"""Generate random inputs for mAP when iou_type=segm."""
633633
preds = []
634634
targets = []
635-
for _ in range(2):
635+
for _ in range(batch_size):
636636
result = {}
637-
num_preds = torch.randint(0, 10, (1,)).item()
637+
num_preds = torch.randint(0, num_preds_size, (1,)).item() if random_size else num_preds_size
638638
result["scores"] = torch.rand((num_preds,), device=device)
639639
result["labels"] = torch.randint(0, 10, (num_preds,), device=device)
640640
result["masks"] = torch.randint(0, 2, (num_preds, 10, 10), device=device).bool()
641641
preds.append(result)
642642
gt = {}
643-
num_gt = torch.randint(0, 10, (1,)).item()
643+
num_gt = torch.randint(0, num_gt_size, (1,)).item() if random_size else num_gt_size
644644
gt["labels"] = torch.randint(0, 10, (num_gt,), device=device)
645645
gt["masks"] = torch.randint(0, 2, (num_gt, 10, 10), device=device).bool()
646646
targets.append(gt)
@@ -690,3 +690,23 @@ def test_for_box_format(box_format, iou_val_expected, map_val_expected):
690690
result = metric.compute()
691691
assert result["map"].item() == map_val_expected
692692
assert round(float(metric.coco_eval.ious[(0, 0)]), 3) == iou_val_expected
693+
694+
695+
@pytest.mark.parametrize("iou_type", ["bbox", "segm"])
696+
def test_warning_on_many_detections(iou_type):
697+
"""Test that a warning is raised when there are many detections."""
698+
if iou_type == "bbox":
699+
preds = [
700+
{
701+
"boxes": torch.tensor([[0.5, 0.5, 1, 1]]).repeat(101, 1),
702+
"scores": torch.tensor([1.0]).repeat(101),
703+
"labels": torch.tensor([0]).repeat(101),
704+
}
705+
]
706+
targets = [{"boxes": torch.tensor([[0, 0, 1, 1]]), "labels": torch.tensor([0])}]
707+
else:
708+
preds, targets = _generate_random_segm_input("cpu", 1, 101, 10, False)
709+
710+
metric = MeanAveragePrecision(iou_type=iou_type)
711+
with pytest.warns(UserWarning, match="Encountered more than 100 detections in a single image.*"):
712+
metric.update(preds, targets)

0 commit comments

Comments
 (0)