torchvision.transforms.GaussianBlur occasionally returns its input unchanged, and doesn't behave deterministically

I was trying to implement a few versions of local image normalization, all involving some variation of a Gaussian blur, then subtracting that from the original image. I kept getting odd results such as occasional images filled with all 0s or all -1s or similar. After some investigation, I was able to narrow it down to a minimal example to reproduce the bug.

It turns out that torchvision.transforms.GaussianBlur will occasionally return a tensor identical to its input, seemingly at random. Not only that, but it doesn’t appear to behave deterministicly, even with torch.use_deterministic_algorithms(True) set.

I’ve made a colab notebook to demonstrate the issue. It initializes an image tensor, clones it 64 times, and runs it through a simplified local normalization function. This function just blurs the image using torchvision.transforms.GaussianBlur, and subtracts the original form it. Not only does the output of this operation vary by several orders of magnitude, even when using deterministic algorithms, the output is occasionally 0, meaning that applying a gaussian(img) == img approximately 2/64 times. This behavior persists on both CPU and GPU.

This very much looks like a bug to me, but perhaps someone who knows Pytorch’s internals better than me can explain something I’m missing.

The code from the notebook is reproduced below:

import torch
import torchvision
import matplotlib
import matplotlib.pyplot as plt
torch.use_deterministic_algorithms(True)

print(torch.__version__)

def localNorm(img):

	img = img.float()
	#img = img.to("cuda") ##Bug exists on both CPU and GPU
	blur = torchvision.transforms.GaussianBlur(9)
	img_b = blur(img)
	
	#assert(not (img_b == img).all()) ##Optional assert to catch the bug here, rather than comparing the sum to 0 later

	img_b = img_b - img ##subtract blurred image from the original

	return img_b.cpu().detach()

##three different ways of initializing the image. Two random, and one deterministic. All exhibit the bug

#r = torch.rand(1,2160,2560)

#r = torch.normal(0,1,(1,2160,2560))

r = torch.linspace(0,2160*2560,2160*2560)#*(3.141592653589)
r = torch.cos(r)
r = torch.unflatten(r,0,(1,2160,2560))

##intit a tensor 'r' representing an image, and clone it 64 times
dataset = [torch.clone(r) for i in range(64)]

fig, axarr = plt.subplots(8,8)

means = []

##run localNorm on the identical image 64 times, and observe differing results, with the occasional identical (zero sum of differences) result
for i,img in enumerate(dataset):
	norm = localNorm(img)[0]
	axarr[i//8,i%8].imshow(norm)#, cmap='gray')
	mean = torch.sum(norm)
	means.append(mean)
	print(mean.item(), (mean==0).item())

plt.gca().set_axis_off()
plt.subplots_adjust(top = 1, bottom = 0, right = 1, left = 0, hspace = 0, wspace = 0)
plt.margins(0,0)
for ax in fig.axes:
	ax.axis('off')

plt.show()

plt.plot(means)
plt.show()

Hi Brian!

By design, an instance of GaussianBlur will use a randomly chosen
value of sigma to generate its gaussian kernel every time it is applied
to an image. This is true regardless of whether or not you have set
torch.use_deterministic_algorithms (True).

When your instance of GaussianBlur chooses a sufficiently small value
for sigma, the kernel it generates can be very close to no blur at all. This
is the effect you are seeing.

Even though you run this in a loop, you instantiate the same “value” of
blur every time the first line runs – there is no randomness yet. Then
every time you run the second line, blur (the GaussianBlur instance)
generates a random value of sigma, uses that sigma to create the kernel,
and uses that kernel to blur the image. Sometimes the generated sigma
is sufficiently small that, up to numerical precision, the image doesn’t get
blurred.

Consider:

>>> import torch
>>> import torchvision
>>>
>>> print (torch.__version__)
2.2.1
>>> print (torchvision.__version__)
0.17.1
>>>
>>> _ = torch.manual_seed (2024)
>>>
>>> delt = torch.zeros (1, 11, 11)                                 # make a "delta-function" image with one central spot
>>> delt[0, 5, 5] = 1.0
>>> delt[0, 3:8, 3:8]                                              # look at just the center of delt
tensor([[0., 0., 0., 0., 0.],
        [0., 0., 0., 0., 0.],
        [0., 0., 1., 0., 0.],
        [0., 0., 0., 0., 0.],
        [0., 0., 0., 0., 0.]])
>>>
>>> blurA = torchvision.transforms.GaussianBlur (9)                # instantiate a GaussianBlur
>>> blurA.sigma                                                    # sigma is the default range of (0.1, 2.0)
(0.1, 2.0)
>>> delt_blurA = blurA (delt)                                      # blur delt with a kernel with a random choice of sigma
>>> delt_blurA[0, 3:8, 3:8]                                        # look at just the center of the kernel
tensor([[0.0050, 0.0170, 0.0255, 0.0170, 0.0050],
        [0.0170, 0.0574, 0.0861, 0.0574, 0.0170],
        [0.0255, 0.0861, 0.1291, 0.0861, 0.0255],
        [0.0170, 0.0574, 0.0861, 0.0574, 0.0170],
        [0.0050, 0.0170, 0.0255, 0.0170, 0.0050]])
>>>
>>> blurB = torchvision.transforms.GaussianBlur (9, sigma = 0.1)   # instantiate a GaussianBlur with fixed sigma
>>> blurB.sigma                                                    # sigma is the fixed value of 0.1
(0.1, 0.1)
>>> delt_blurB = blurB (delt)                                      # blur delt with sigma = 0.1
>>> delt_blurB[0, 3:8, 3:8]                                        # kernel is essentially no blur
tensor([[0.0000e+00, 0.0000e+00, 0.0000e+00, 0.0000e+00, 0.0000e+00],
        [0.0000e+00, 3.7835e-44, 1.9287e-22, 3.7835e-44, 0.0000e+00],
        [0.0000e+00, 1.9287e-22, 1.0000e+00, 1.9287e-22, 0.0000e+00],
        [0.0000e+00, 3.7835e-44, 1.9287e-22, 3.7835e-44, 0.0000e+00],
        [0.0000e+00, 0.0000e+00, 0.0000e+00, 0.0000e+00, 0.0000e+00]])

Best.

K. Frank