Skip to content

Commit

Permalink
Merges bbox overlap implementations and adds tests
Browse files Browse the repository at this point in the history
  • Loading branch information
vierja committed Aug 11, 2017
1 parent 51afb2d commit c46d18f
Show file tree
Hide file tree
Showing 6 changed files with 149 additions and 74 deletions.
4 changes: 2 additions & 2 deletions luminoth/models/fasterrcnn/rcnn_target.py
Original file line number Diff line number Diff line change
Expand Up @@ -3,7 +3,7 @@
import numpy as np

from luminoth.utils.bbox_transform import encode, unmap
from luminoth.utils.bbox import bbox_overlaps
from luminoth.utils.bbox_overlap import bbox_overlap


class RCNNTarget(snt.AbstractModule):
Expand Down Expand Up @@ -95,7 +95,7 @@ def proposal_target_layer(self, proposals, gt_boxes):
# Remove batch id from proposals
proposals = proposals[:, 1:]

overlaps = bbox_overlaps(
overlaps = bbox_overlap(
# We need to use float and ascontiguousarray because of Cython
# implementation of bbox_overlap
np.ascontiguousarray(proposals, dtype=np.float),
Expand Down
4 changes: 2 additions & 2 deletions luminoth/models/fasterrcnn/rpn_anchor_target.py
Original file line number Diff line number Diff line change
Expand Up @@ -2,7 +2,7 @@
import tensorflow as tf
import numpy as np

from luminoth.utils.bbox import bbox_overlaps
from luminoth.utils.bbox_overlap import bbox_overlap
from luminoth.utils.bbox_transform import encode, unmap


Expand Down Expand Up @@ -159,7 +159,7 @@ def _anchor_target_layer_np(self, pretrained_shape, gt_boxes, im_size, all_ancho

# intersection over union (IoU) overlap between the anchors and the
# ground truth boxes.
overlaps = bbox_overlaps(
overlaps = bbox_overlap(
np.ascontiguousarray(anchors, dtype=np.float),
np.ascontiguousarray(gt_boxes, dtype=np.float))

Expand Down
47 changes: 0 additions & 47 deletions luminoth/utils/bbox.py

This file was deleted.

75 changes: 56 additions & 19 deletions luminoth/utils/bbox_overlap.py
Original file line number Diff line number Diff line change
@@ -1,8 +1,11 @@
import numpy as np
import tensorflow as tf


def bbox_overlap_iou(bboxes1, bboxes2):
"""
def bbox_overlap_tf(bboxes1, bboxes2):
"""Calculate Intersection over Union (IoU) between two sets of bounding
boxes.
Args:
bboxes1: shape (total_bboxes1, 4)
with x1, y1, x2, y2 point order.
Expand All @@ -17,9 +20,6 @@ def bbox_overlap_iou(bboxes1, bboxes2):
Tensor with shape (total_bboxes1, total_bboxes2)
with the IoU (intersection over union) of bboxes1[i] and bboxes2[j]
in [i, j].
TODO: Check for invalid values when union is zero or intersection is negative?
"""
with tf.name_scope('bbox_overlap'):
x11, y11, x12, y12 = tf.split(bboxes1, 4, axis=1)
Expand All @@ -31,25 +31,62 @@ def bbox_overlap_iou(bboxes1, bboxes2):
xI2 = tf.minimum(x12, tf.transpose(x22))
yI2 = tf.minimum(y12, tf.transpose(y22))

inter_area = (xI2 - xI1 + 1) * (yI2 - yI1 + 1)
intersection = (
tf.maximum(xI2 - xI1 + 1., 0.) *
tf.maximum(yI2 - yI1 + 1., 0.)
)

bboxes1_area = (x12 - x11 + 1) * (y12 - y11 + 1)
bboxes2_area = (x22 - x21 + 1) * (y22 - y21 + 1)

union = (bboxes1_area + tf.transpose(bboxes2_area)) - inter_area
union = (bboxes1_area + tf.transpose(bboxes2_area)) - intersection

iou = tf.maximum(intersection / union, 0)

return iou


def bbox_overlap(bboxes1, bboxes2):
"""Calculate Intersection of Union between two sets of bounding boxes.
Intersection over Union (IoU) of two bounding boxes A and B is calculated
doing: (A ∩ B) / (A ∪ B).
Args:
bboxes1: numpy array of shape (total_bboxes1, 4).
bboxes2: numpy array of shape (total_bboxes2, 4).
Returns:
iou: numpy array of shape (total_bboxes1, total_bboxes1) a matrix with
the intersection over union of bboxes1[i] and bboxes2[j] in
iou[i][j].
"""
xI1 = np.maximum(bboxes1[:, [0]], bboxes2[:, [0]].T)
yI1 = np.maximum(bboxes1[:, [1]], bboxes2[:, [1]].T)

xI2 = np.minimum(bboxes1[:, [2]], bboxes2[:, [2]].T)
yI2 = np.minimum(bboxes1[:, [3]], bboxes2[:, [3]].T)

intersection = (
np.maximum(xI2 - xI1 + 1, 0.) *
np.maximum(yI2 - yI1 + 1, 0.)
)

return tf.maximum(inter_area / union, 0)
bboxes1_area = (
(bboxes1[:, [2]] - bboxes1[:, [0]] + 1) *
(bboxes1[:, [3]] - bboxes1[:, [1]] + 1)
)
bboxes2_area = (
(bboxes2[:, [2]] - bboxes2[:, [0]] + 1) *
(bboxes2[:, [3]] - bboxes2[:, [1]] + 1)
)

# Calculate the union as the sum of areas minus intersection
union = (bboxes1_area + bboxes2_area.T) - intersection

if __name__ == '__main__':
bboxes1 = tf.placeholder(tf.float32)
bboxes2 = tf.placeholder(tf.float32)
overlap_op = bbox_overlap_iou(bboxes1, bboxes2)
# We start we an empty array of zeros.
iou = np.zeros((bboxes1.shape[0], bboxes2.shape[0]))

bboxes1_vals = [[39, 63, 203, 112], [0, 0, 10, 10]]
bboxes2_vals = [[3, 4, 24, 32], [54, 66, 198, 114], [6, 7, 60, 44]]
with tf.Session() as sess:
overlap = sess.run(overlap_op, feed_dict={
bboxes1: bboxes1_vals,
bboxes2: bboxes2_vals,
})
# Only divide where the intersection is > 0
np.divide(intersection, union, out=iou, where=intersection > 0.)
return iou
85 changes: 85 additions & 0 deletions luminoth/utils/bbox_overlap_test.py
Original file line number Diff line number Diff line change
@@ -0,0 +1,85 @@
import numpy as np
import tensorflow as tf

from luminoth.utils.bbox_overlap import bbox_overlap_tf, bbox_overlap


class BBoxOverlapTest(tf.test.TestCase):
"""Tests for bbox_overlap
bbox_overlap has a TensorFlow and a Numpy implementation.
We test both at the same time by getting both values and making sure they
are both equal before doing any assertions.
"""
def _get_iou(self, bbox1_val, bbox2_val):
"""Get IoU for two sets of bounding boxes.
It also checks that both implementations return the same before
returning.
Args:
bbox1_val: Array of shape (total_bbox1, 4).
bbox2_val: Array of shape (total_bbox2, 4).
Returns:
iou: Array of shape (total_bbox1, total_bbox2)
"""
bbox1 = tf.placeholder(tf.float32, (None, 4))
bbox2 = tf.placeholder(tf.float32, (None, 4))
iou = bbox_overlap_tf(bbox1, bbox2)

with self.test_session() as sess:
iou_val_tf = sess.run(iou, feed_dict={
bbox1: np.array(bbox1_val),
bbox2: np.array(bbox2_val),
})

iou_val_np = bbox_overlap(np.array(bbox1_val), np.array(bbox2_val))
self.assertAllClose(iou_val_np, iou_val_tf)
return iou_val_tf

def testNoOverlap(self):
# Single box test
iou = self._get_iou([[0, 0, 10, 10]], [[11, 11, 20, 20]])
self.assertAllEqual(iou, [[0.]])

# Multiple boxes.
iou = self._get_iou(
[[0, 0, 10, 10], [5, 5, 10, 10]],
[[11, 11, 20, 20], [15, 15, 20, 20]]
)
self.assertAllEqual(iou, [[0., 0.], [0., 0.]])

def testAllOverlap(self):
# Equal boxes
iou = self._get_iou([[0, 0, 10, 10]], [[0, 0, 10, 10]])
self.assertAllEqual(iou, [[1.]])

# Crossed equal boxes.
iou = self._get_iou(
[[0, 0, 10, 10], [11, 11, 20, 20]],
[[0, 0, 10, 10], [11, 11, 20, 20]]
)
# We should get an identity matrix.
self.assertAllEqual(iou, [[1., 0.], [0., 1.]])

def testInvalidBoxes(self):
# Zero area, bbox1 has x_min == x_max
iou = self._get_iou([[10, 0, 10, 10]], [[0, 0, 10, 10]])
# self.assertAllEqual(iou, [[0.]]) TODO: Fails

# Negative area, bbox1 has x_min > x_max (only by one)
iou = self._get_iou([[10, 0, 9, 10]], [[0, 0, 10, 10]])
self.assertAllEqual(iou, [[0.]])

# Negative area, bbox1 has x_min > x_max
iou = self._get_iou([[10, 0, 7, 10]], [[0, 0, 10, 10]])
self.assertAllEqual(iou, [[0.]])

# Negative area in both cases, both boxes equal but negative
iou = self._get_iou([[10, 0, 7, 10]], [[10, 0, 7, 10]])
self.assertAllEqual(iou, [[0.]])


if __name__ == '__main__':
tf.test.main()
8 changes: 4 additions & 4 deletions luminoth/utils/image_vis.py
Original file line number Diff line number Diff line change
Expand Up @@ -6,7 +6,7 @@
import PIL.ImageFont as ImageFont
import tensorflow as tf

from .bbox import bbox_overlaps
from .bbox_overlap import bbox_overlap
from .bbox_transform import decode
from base64 import b64encode
from sys import stdout
Expand Down Expand Up @@ -739,10 +739,10 @@ def draw_object_prediction(pred_dict, topn=50):

def draw_rcnn_input_proposals(pred_dict):
tf.logging.debug('Display RPN proposals used in training classification. Top IoU with GT is displayed.')
proposals = pred_dict['rpn_prediction']['proposals'][:,1:]
gt_boxes = pred_dict['gt_boxes'][:,:4]
proposals = pred_dict['rpn_prediction']['proposals'][:, 1:]
gt_boxes = pred_dict['gt_boxes'][:, :4]

overlaps = bbox_overlaps(
overlaps = bbox_overlap(
np.ascontiguousarray(proposals, dtype=np.float),
np.ascontiguousarray(gt_boxes, dtype=np.float)
)
Expand Down

0 comments on commit c46d18f

Please sign in to comment.