Skip to content

Commit fa56090

Browse files
authored
Merge branch 'dev' into fix-issue-8780
Signed-off-by: Eric Kerfoot <17726042+ericspod@users.noreply.github.com>
2 parents f2fa381 + 07990cc commit fa56090

3 files changed

Lines changed: 70 additions & 9 deletions

File tree

monai/losses/image_dissimilarity.py

Lines changed: 6 additions & 6 deletions
Original file line numberDiff line numberDiff line change
@@ -15,7 +15,7 @@
1515
from torch.nn import functional as F
1616
from torch.nn.modules.loss import _Loss
1717

18-
from monai.networks.layers import gaussian_1d, separable_filtering
18+
from monai.networks.layers import separable_filtering
1919
from monai.utils import LossReduction
2020
from monai.utils.module import look_up_option
2121

@@ -34,11 +34,11 @@ def make_triangular_kernel(kernel_size: int) -> torch.Tensor:
3434

3535

3636
def make_gaussian_kernel(kernel_size: int) -> torch.Tensor:
37-
sigma = torch.tensor(kernel_size / 3.0)
38-
kernel = gaussian_1d(sigma=sigma, truncated=(kernel_size // 2) / sigma, approx="sampled", normalize=False) * (
39-
2.5066282 * sigma
40-
)
41-
return kernel[:kernel_size]
37+
sigma = kernel_size / 3.0
38+
half = kernel_size // 2
39+
x = torch.arange(-half, half + 1, dtype=torch.float)
40+
kernel = torch.exp(-0.5 / (sigma * sigma) * x**2)
41+
return kernel
4242

4343

4444
kernel_dict = {

tests/data/utils/test_compute_shape_offset.py

Lines changed: 35 additions & 2 deletions
Original file line numberDiff line numberDiff line change
@@ -19,8 +19,8 @@
1919
from monai.data.utils import compute_shape_offset
2020

2121

22-
class TestComputeShapeOffsetRegression(unittest.TestCase):
23-
"""Regression tests for `compute_shape_offset` input-shape handling."""
22+
class TestComputeShapeOffset(unittest.TestCase):
23+
"""Unit tests for :func:`monai.data.utils.compute_shape_offset`."""
2424

2525
def test_pytorch_size_input(self):
2626
"""Validate `torch.Size` input produces expected shape and offset.
@@ -42,6 +42,39 @@ def test_pytorch_size_input(self):
4242
# 3. Prove it successfully processed the shape by checking its length
4343
self.assertEqual(len(shape), 3)
4444

45+
def setUp(self):
46+
"""Set up a 4x4 identity affine used across all test cases."""
47+
self.affine = np.eye(4)
48+
49+
def test_numpy_array_input(self):
50+
"""Verify compute_shape_offset accepts a numpy array as spatial_shape."""
51+
shape = np.array([64, 64, 64])
52+
out_shape, _ = compute_shape_offset(shape, self.affine, self.affine)
53+
self.assertEqual(len(out_shape), 3)
54+
55+
def test_list_input(self):
56+
"""Verify compute_shape_offset accepts a plain list as spatial_shape."""
57+
shape = [64, 64, 64]
58+
out_shape, _ = compute_shape_offset(shape, self.affine, self.affine)
59+
self.assertEqual(len(out_shape), 3)
60+
61+
def test_torch_tensor_input(self):
62+
"""Verify compute_shape_offset accepts a torch.Tensor as spatial_shape.
63+
64+
This path broke in PyTorch >= 2.9 because np.array() relied on the
65+
non-tuple sequence indexing protocol that PyTorch removed. Wrapping with
66+
tuple() fixes it.
67+
"""
68+
shape = torch.tensor([64, 64, 64])
69+
out_shape, _ = compute_shape_offset(shape, self.affine, self.affine)
70+
self.assertEqual(len(out_shape), 3)
71+
72+
def test_identity_affines_preserve_shape(self):
73+
"""Verify that identity in/out affines produce an output shape matching the input."""
74+
shape = torch.tensor([32, 48, 16])
75+
out_shape, _ = compute_shape_offset(shape, self.affine, self.affine)
76+
np.testing.assert_allclose(np.array(out_shape, dtype=float), shape.numpy().astype(float), atol=1e-5)
77+
4578

4679
if __name__ == "__main__":
4780
unittest.main()

tests/losses/image_dissimilarity/test_local_normalized_cross_correlation_loss.py

Lines changed: 29 additions & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -17,7 +17,7 @@
1717
import torch
1818
from parameterized import parameterized
1919

20-
from monai.losses.image_dissimilarity import LocalNormalizedCrossCorrelationLoss
20+
from monai.losses.image_dissimilarity import LocalNormalizedCrossCorrelationLoss, make_gaussian_kernel
2121

2222
device = "cuda" if torch.cuda.is_available() else "cpu"
2323

@@ -113,6 +113,25 @@
113113
},
114114
-0.95406944,
115115
],
116+
# Regression tests for gh-8780: gaussian kernel_size > 3 was broken due to
117+
# truncated parameter being passed as pixel radius instead of sigma multiplier.
118+
# Identical images must yield loss == -1.0 for any kernel size.
119+
[
120+
{"spatial_dims": 1, "kernel_type": "gaussian", "kernel_size": 5},
121+
{
122+
"pred": torch.arange(0, 5).reshape(1, 1, -1).to(dtype=torch.float, device=device),
123+
"target": torch.arange(0, 5).reshape(1, 1, -1).to(dtype=torch.float, device=device),
124+
},
125+
-1.0,
126+
],
127+
[
128+
{"spatial_dims": 1, "kernel_type": "gaussian", "kernel_size": 9},
129+
{
130+
"pred": torch.arange(0, 9).reshape(1, 1, -1).to(dtype=torch.float, device=device),
131+
"target": torch.arange(0, 9).reshape(1, 1, -1).to(dtype=torch.float, device=device),
132+
},
133+
-1.0,
134+
],
116135
]
117136

118137

@@ -138,6 +157,15 @@ def test_ill_shape(self):
138157
torch.ones((1, 3, 4, 4, 4), dtype=torch.float, device=device),
139158
)
140159

160+
def test_gaussian_kernel_shape_and_symmetry(self):
161+
# gh-8780: kernel must have correct length, be symmetric, and peak at center
162+
for kernel_size in [3, 5, 7, 9, 11, 15]:
163+
k = make_gaussian_kernel(kernel_size)
164+
self.assertEqual(len(k), kernel_size)
165+
self.assertTrue(torch.allclose(k, k.flip(0)), f"kernel_size={kernel_size} not symmetric")
166+
self.assertEqual(k.argmax().item(), kernel_size // 2)
167+
np.testing.assert_allclose(k.max().item(), 1.0, rtol=1e-6)
168+
141169
def test_ill_opts(self):
142170
pred = torch.ones((1, 3, 3, 3, 3), dtype=torch.float)
143171
target = torch.ones((1, 3, 3, 3, 3), dtype=torch.float)

0 commit comments

Comments
 (0)